Skip to content

Commit bf36620

Browse files
authored
[Bug fixes]: Modify code and setting to run ACT correctly (#3539)
* Modify the code in test_seg.py to collect dynamic shape information for dataset ADE20k * 调整ppmobileseg的验证配置,更新全部测试结果以及环境安装文档 * modify
1 parent 3c3a12a commit bf36620

File tree

4 files changed

+93
-32
lines changed

4 files changed

+93
-32
lines changed

deploy/slim/act/configs/datasets/ade_data.yml

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -22,6 +22,10 @@ train_dataset:
2222

2323
val_dataset:
2424
transforms:
25+
- type: Resize
26+
target_size: [2048, 512]
27+
keep_ratio: True
28+
size_divisor: 32
2529
- type: Normalize
2630
mean: [0.485, 0.456, 0.406]
2731
std: [0.229, 0.224, 0.225]

deploy/slim/act/readme.md

Lines changed: 21 additions & 20 deletions
Original file line numberDiff line numberDiff line change
@@ -17,30 +17,30 @@
1717

1818
## 2.Benchmark
1919

20-
| 模型 | 策略 | Total IoU (%) | CPU耗时(ms)<br>thread=10<br>mkldnn=on| Nvidia GPU耗时(ms)<br>TRT=on| 配置文件 | Inference模型 |
20+
| 模型 | 策略 | Total IoU (%) | CPU耗时(ms)<br>thread=12<br>mkldnn=on | Nvidia GPU耗时(ms)<br>TRT=on| 配置文件 | Inference模型 |
2121
|:-----:|:-----:|:----------:|:---------:| :------:|:------:|:------:|
22-
| OCRNet_HRNetW48 |Baseline |82.15| **4332.2** | **154.9** | - | [mode](https://paddleseg.bj.bcebos.com/deploy/slim_act/ocrnet/ocrnet_export.zip)|
23-
| OCRNet_HRNetW48 | 量化蒸馏训练 |82.03| **3728.7** | **59.8**|[config](configs/ocrnet/ocrnet_hrnetw48_qat.yaml)| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ocrnet/ocrnet_qat.zip) |
24-
| SegFormer-B0* |Baseline | 75.27| 285.4| 34.3 |-| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/segformer/segformer_b0_export.zip) |
25-
| SegFormer-B0* |量化蒸馏训练 | 75.22 | 284.1| 35.7|[config](configs/segformer/segformer_b0_qat.yaml)| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/segformer/segformer_qat.zip) |
26-
| PP-LiteSeg-Tiny |Baseline | 77.04 | 640.72 | **11.9** | - |[model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppliteseg/liteseg_tiny_scale1.0.zip)|
27-
| PP-LiteSeg-Tiny |量化蒸馏训练 | 77.14 | 450.19 | **7.5** | [config](./configs/ppliteseg/ppliteseg_qat.yaml)|[model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppliteseg/save_quant_model_qat.zip)|
28-
| PP-MobileSeg-Base |Baseline |41.55| **311.1** | **17.8** | - | [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppmobileseg/ppmobileseg_base_ade_export.zip) |
29-
| PP-MobileSeg-Base |量化蒸馏训练 |39.08| **303.6** | **16.2**| [config](configs/ppmobileseg/ppmobileseg_qat.yml)| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppmobileseg/ppmobileseg_base_ade.zip)|
30-
31-
* SegFormer-B0 is tested on CPU under deleted gpu_cpu_map_matmul_v2_to_mul_pass because it will raise an error.
22+
| OCRNet_HRNetW48 |Baseline |82.16| **5788.7** | **153.0** | - | [mode](https://paddleseg.bj.bcebos.com/deploy/slim_act/ocrnet/ocrnet_export.zip)|
23+
| OCRNet_HRNetW48 | 量化蒸馏训练 |82.02| **5291.4** | **60.0** |[config](configs/ocrnet/ocrnet_hrnetw48_qat.yaml)| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ocrnet/ocrnet_qat.zip) |
24+
| SegFormer-B0* |Baseline | 75.27| **3234.6** | **72.6** |-| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/segformer/segformer_b0_export.zip) |
25+
| SegFormer-B0* |量化蒸馏训练 | 75.26 | **2906.2** | **52.4** |[config](configs/segformer/segformer_b0_qat.yaml)| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/segformer/segformer_qat.zip) |
26+
| PP-LiteSeg-Tiny |Baseline | 77.04 | 1038.4 | **11.7** | - |[model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppliteseg/liteseg_tiny_scale1.0.zip)|
27+
| PP-LiteSeg-Tiny |量化蒸馏训练 | 77.16 | 1163.8 | **7.2** | [config](./configs/ppliteseg/ppliteseg_qat.yaml)|[model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppliteseg/save_quant_model_qat.zip)|
28+
| PP-MobileSeg-Base |Baseline |40.69| **547.7** | **22.3** | - | [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppmobileseg/ppmobileseg_base_ade_export.zip) |
29+
| PP-MobileSeg-Base |量化蒸馏训练 |38.18| **439.8** | **21.1** | [config](configs/ppmobileseg/ppmobileseg_qat.yml)| [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppmobileseg/ppmobileseg_base_ade.zip)|
30+
| PP-MobileSeg-Base |量化蒸馏训练(关闭IR优化) |39.92| **1296.3** | **44.3** | - | [model](https://paddleseg.bj.bcebos.com/deploy/slim_act/ppmobileseg/ppmobileseg_base_ade.zip)|
31+
3232
* PP-MobileSeg-Base is tested on ADE20K dataset, while others are tested on cityscapes.
3333

3434
- CPU测试环境:
35-
- Intel(R) Xeon(R) Gold 6271C CPU @ 2.60GHz
35+
- Intel(R) Xeon(R) Gold 6148 CPU @ 2.40GHz
3636
- cpu thread: 10
3737

3838

3939
- Nvidia GPU测试环境:
4040

4141
- 硬件:NVIDIA Tesla V100 单卡
4242
- 软件:CUDA 11.2, cudnn 8.1.0, TensorRT-8.0.3.4
43-
- 测试配置:batch_size: 4
43+
- 测试配置:batch_size: 1
4444

4545
- 测速要求:
4646
- 批量测试取平均:单张图片上测速时间会有浮动,因此测速需要跑10遍warmup,再跑100次取平均。现有test_seg的批量测试已经集成该功能。
@@ -53,21 +53,22 @@
5353

5454
#### 3.1 准备环境
5555

56-
- PaddlePaddle == 2.5 (可从[Paddle官网](https://www.paddlepaddle.org.cn/install/quick?docurl=/documentation/docs/zh/install/pip/linux-pip.html)下载安装)
57-
- PaddleSlim == 2.5
56+
- PaddlePaddle == develop (可从[Paddle官网](https://www.paddlepaddle.org.cn/install/quick?docurl=/documentation/docs/zh/install/pip/linux-pip.html)下载安装)
57+
- PaddleSlim == develop
5858
- PaddleSeg == develop
5959

6060
安装paddlepaddle:
6161
```shell
6262
# CPU
63-
python -m pip install paddlepaddle==2.5.1 -i https://pypi.tuna.tsinghua.edu.cn/simple
64-
# GPU 以Ubuntu、CUDA 10.2为例
65-
python -m pip install paddlepaddle-gpu==2.5.1.post102 -f https://www.paddlepaddle.org.cn/whl/linux/mkl/avx/stable.html
63+
python -m pip install paddlepaddle==0.0.0 -f https://www.paddlepaddle.org.cn/whl/linux/cpu-mkl/develop.html
64+
# GPU 以Ubuntu、CUDA 11.2为例
65+
python -m pip install paddlepaddle-gpu==0.0.0.post112 -f https://www.paddlepaddle.org.cn/whl/linux/gpu/develop.html
6666
```
6767

68-
安装paddleslim 2.5
68+
安装paddleslim develop
6969
```shell
70-
pip install paddleslim@git+https://gitee.com/paddlepaddle/PaddleSlim.git@release/2.5
70+
git clone https://github.com/PaddlePaddle/PaddleSlim.git & cd PaddleSlim
71+
python setup.py install
7172
```
7273

7374
安装paddleseg develop和对应包:

deploy/slim/act/run_seg.py

Lines changed: 4 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -96,9 +96,11 @@ def eval_function(exe, compiled_test_program, test_feed_names, test_fetch_list):
9696
paddle.disable_static()
9797
logit = logits[
9898
0] # logit shape is 3, except data['trans_info'] needs to be empty
99+
for i in range(len(data['trans_info'][::-1][0][1])):
100+
data['trans_info'][::-1][0][1][i] = paddle.to_tensor(data['trans_info'][::-1][0][1][i])
99101
logit = reverse_transform(
100-
paddle.to_tensor(logit), data['trans_info'], mode='bilinear')
101-
pred = paddle.to_tensor(logit)
102+
paddle.to_tensor(logit).unsqueeze(0), data['trans_info'], mode='bilinear')
103+
pred = paddle.to_tensor(logit).squeeze(0)
102104
if len(
103105
pred.shape
104106
) == 4: # for humanseg model whose prediction is distribution but not class id

deploy/slim/act/test_seg.py

Lines changed: 64 additions & 10 deletions
Original file line numberDiff line numberDiff line change
@@ -41,6 +41,39 @@ def _transforms(dataset):
4141
return transforms
4242

4343

44+
def find_images_with_bounding_size(eval_dataset: paddle.io.Dataset):
45+
max_length_index = -1
46+
max_width_index = -1
47+
min_length_index = -1
48+
min_width_index = -1
49+
50+
max_length = float('-inf')
51+
max_width = float('-inf')
52+
min_length = float('inf')
53+
min_width = float('inf')
54+
for idx, data in enumerate(eval_dataset):
55+
image = np.array(data['img'])
56+
h, w = image.shape[-2:]
57+
if h > max_length:
58+
max_length = h
59+
max_length_index = idx
60+
if w > max_width:
61+
max_width = w
62+
max_width_index = idx
63+
if h < min_length:
64+
min_length = h
65+
min_length_index = idx
66+
if w < min_width:
67+
min_width = w
68+
min_width_index = idx
69+
print(f"Found max image length: {max_length}, index: {max_length_index}")
70+
print(f"Found max image width: {max_width}, index: {max_width_index}")
71+
print(f"Found min image length: {min_length}, index: {min_length_index}")
72+
print(f"Found min image width: {min_width}, index: {min_width_index}")
73+
return paddle.io.Subset(eval_dataset, [max_width_index, max_length_index,
74+
min_width_index, min_length_index])
75+
76+
4477
def load_predictor(args):
4578
"""
4679
load predictor func
@@ -109,7 +142,7 @@ def predict_image(args):
109142
data = transform({'img': args.image_file})
110143
data = data['img'][np.newaxis, :]
111144

112-
# Step2: Prepare prdictor
145+
# Step2: Prepare predictor
113146
predictor, rerun_flag = load_predictor(args)
114147

115148
# Step3: Inference
@@ -167,6 +200,15 @@ def eval(args):
167200

168201
eval_dataset = builder.val_dataset
169202

203+
predictor, rerun_flag = load_predictor(args)
204+
205+
if rerun_flag and args.use_multi_img_for_dynamic_shape_collect:
206+
print(
207+
"***** Try to find the images with the largest and smallest length and width respectively in the ADE20K "
208+
"dataset for collecting dynamic shape. *****"
209+
)
210+
eval_dataset = find_images_with_bounding_size(eval_dataset)
211+
170212
batch_sampler = paddle.io.BatchSampler(
171213
eval_dataset, batch_size=1, shuffle=False, drop_last=False)
172214
loader = paddle.io.DataLoader(
@@ -175,8 +217,6 @@ def eval(args):
175217
num_workers=0,
176218
return_list=True)
177219

178-
predictor, rerun_flag = load_predictor(args)
179-
180220
intersect_area_all = 0
181221
pred_area_all = 0
182222
label_area_all = 0
@@ -207,14 +247,22 @@ def eval(args):
207247
time_max = max(time_max, timed)
208248
predict_time += timed
209249
if rerun_flag:
210-
print(
211-
"***** Collect dynamic shape done, Please rerun the program to get correct results. *****"
212-
)
213-
return
214-
250+
if args.use_multi_img_for_dynamic_shape_collect:
251+
if batch_id == sample_nums - 1:
252+
print(
253+
"***** Collect dynamic shape done, Please rerun the program to get correct results. *****"
254+
)
255+
return
256+
else:
257+
continue
258+
else:
259+
print(
260+
"***** Collect dynamic shape done, Please rerun the program to get correct results. *****"
261+
)
262+
return
215263
logit = reverse_transform(
216-
paddle.to_tensor(results), data['trans_info'], mode="bilinear")
217-
pred = paddle.to_tensor(logit)
264+
paddle.to_tensor(results).unsqueeze(0), data['trans_info'], mode="bilinear")
265+
pred = paddle.to_tensor(logit).squeeze(0)
218266
if len(
219267
pred.shape
220268
) == 4: # for humanseg model whose prediction is distribution but not class id
@@ -314,6 +362,12 @@ def eval(args):
314362
help="Whether use mkldnn or not.")
315363
parser.add_argument(
316364
"--cpu_threads", type=int, default=1, help="Num of cpu threads.")
365+
parser.add_argument(
366+
"--use_multi_img_for_dynamic_shape_collect",
367+
type=bool,
368+
default=False,
369+
help="Whether it is necessary to use multiple images to collect shape infomation,\
370+
When the image sizes in the data set are different, it needs to be set to True.")
317371
args = parser.parse_args()
318372
if args.image_file:
319373
predict_image(args)

0 commit comments

Comments
 (0)