Skip to content

Commit 769c810

Browse files
Morty-XuhhaAndroid
andauthored
[CodeCamp2023-504]Add a new script to support the WBF
Co-authored-by: huanghaian <[email protected]>
1 parent 59b0fc5 commit 769c810

File tree

6 files changed

+754
-1
lines changed

6 files changed

+754
-1
lines changed

demo/demo_multi_model.py

Lines changed: 212 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,212 @@
1+
# Copyright (c) OpenMMLab. All rights reserved.
2+
"""Support for multi-model fusion, and currently only the Weighted Box Fusion
3+
(WBF) fusion method is supported.
4+
5+
References: https://github.com/ZFTurbo/Weighted-Boxes-Fusion
6+
7+
Example:
8+
9+
python demo/demo_multi_model.py demo/demo.jpg \
10+
./configs/faster_rcnn/faster-rcnn_r50-caffe_fpn_1x_coco.py \
11+
./configs/retinanet/retinanet_r50-caffe_fpn_1x_coco.py \
12+
--checkpoints \
13+
https://download.openmmlab.com/mmdetection/v2.0/faster_rcnn/faster_rcnn_r50_caffe_fpn_1x_coco/faster_rcnn_r50_caffe_fpn_1x_coco_bbox_mAP-0.378_20200504_180032-c5925ee5.pth \ # noqa
14+
https://download.openmmlab.com/mmdetection/v2.0/retinanet/retinanet_r50_caffe_fpn_1x_coco/retinanet_r50_caffe_fpn_1x_coco_20200531-f11027c5.pth \
15+
--weights 1 2
16+
"""
17+
18+
import argparse
19+
import os.path as osp
20+
21+
import mmcv
22+
import mmengine
23+
from mmengine.fileio import isdir, join_path, list_dir_or_file
24+
from mmengine.logging import print_log
25+
from mmengine.structures import InstanceData
26+
27+
from mmdet.apis import DetInferencer
28+
from mmdet.models.utils import weighted_boxes_fusion
29+
from mmdet.registry import VISUALIZERS
30+
from mmdet.structures import DetDataSample
31+
32+
IMG_EXTENSIONS = ('.jpg', '.jpeg', '.png', '.ppm', '.bmp', '.pgm', '.tif',
33+
'.tiff', '.webp')
34+
35+
36+
def parse_args():
37+
parser = argparse.ArgumentParser(
38+
description='MMDetection multi-model inference demo')
39+
parser.add_argument(
40+
'inputs', type=str, help='Input image file or folder path.')
41+
parser.add_argument(
42+
'config',
43+
type=str,
44+
nargs='*',
45+
help='Config file(s), support receive multiple files')
46+
parser.add_argument(
47+
'--checkpoints',
48+
type=str,
49+
nargs='*',
50+
help='Checkpoint file(s), support receive multiple files, '
51+
'remember to correspond to the above config',
52+
)
53+
parser.add_argument(
54+
'--weights',
55+
type=float,
56+
nargs='*',
57+
default=None,
58+
help='weights for each model, remember to '
59+
'correspond to the above config')
60+
parser.add_argument(
61+
'--fusion-iou-thr',
62+
type=float,
63+
default=0.55,
64+
help='IoU value for boxes to be a match in wbf')
65+
parser.add_argument(
66+
'--skip-box-thr',
67+
type=float,
68+
default=0.0,
69+
help='exclude boxes with score lower than this variable in wbf')
70+
parser.add_argument(
71+
'--conf-type',
72+
type=str,
73+
default='avg', # avg, max, box_and_model_avg, absent_model_aware_avg
74+
help='how to calculate confidence in weighted boxes in wbf')
75+
parser.add_argument(
76+
'--out-dir',
77+
type=str,
78+
default='outputs',
79+
help='Output directory of images or prediction results.')
80+
parser.add_argument(
81+
'--device', default='cuda:0', help='Device used for inference')
82+
parser.add_argument(
83+
'--pred-score-thr',
84+
type=float,
85+
default=0.3,
86+
help='bbox score threshold')
87+
parser.add_argument(
88+
'--batch-size', type=int, default=1, help='Inference batch size.')
89+
parser.add_argument(
90+
'--show',
91+
action='store_true',
92+
help='Display the image in a popup window.')
93+
parser.add_argument(
94+
'--no-save-vis',
95+
action='store_true',
96+
help='Do not save detection vis results')
97+
parser.add_argument(
98+
'--no-save-pred',
99+
action='store_true',
100+
help='Do not save detection json results')
101+
parser.add_argument(
102+
'--palette',
103+
default='none',
104+
choices=['coco', 'voc', 'citys', 'random', 'none'],
105+
help='Color palette used for visualization')
106+
107+
args = parser.parse_args()
108+
109+
if args.no_save_vis and args.no_save_pred:
110+
args.out_dir = ''
111+
112+
return args
113+
114+
115+
def main():
116+
args = parse_args()
117+
118+
results = []
119+
cfg_visualizer = None
120+
dataset_meta = None
121+
122+
inputs = []
123+
filename_list = []
124+
if isdir(args.inputs):
125+
dir = list_dir_or_file(
126+
args.inputs, list_dir=False, suffix=IMG_EXTENSIONS)
127+
for filename in dir:
128+
img = mmcv.imread(join_path(args.inputs, filename))
129+
inputs.append(img)
130+
filename_list.append(filename)
131+
else:
132+
img = mmcv.imread(args.inputs)
133+
inputs.append(img)
134+
img_name = osp.basename(args.inputs)
135+
filename_list.append(img_name)
136+
137+
for i, (config,
138+
checkpoint) in enumerate(zip(args.config, args.checkpoints)):
139+
inferencer = DetInferencer(
140+
config, checkpoint, device=args.device, palette=args.palette)
141+
142+
result_raw = inferencer(
143+
inputs=inputs,
144+
batch_size=args.batch_size,
145+
no_save_vis=True,
146+
pred_score_thr=args.pred_score_thr)
147+
148+
if i == 0:
149+
cfg_visualizer = inferencer.cfg.visualizer
150+
dataset_meta = inferencer.model.dataset_meta
151+
results = [{
152+
'bboxes_list': [],
153+
'scores_list': [],
154+
'labels_list': []
155+
} for _ in range(len(result_raw['predictions']))]
156+
157+
for res, raw in zip(results, result_raw['predictions']):
158+
res['bboxes_list'].append(raw['bboxes'])
159+
res['scores_list'].append(raw['scores'])
160+
res['labels_list'].append(raw['labels'])
161+
162+
visualizer = VISUALIZERS.build(cfg_visualizer)
163+
visualizer.dataset_meta = dataset_meta
164+
165+
for i in range(len(results)):
166+
bboxes, scores, labels = weighted_boxes_fusion(
167+
results[i]['bboxes_list'],
168+
results[i]['scores_list'],
169+
results[i]['labels_list'],
170+
weights=args.weights,
171+
iou_thr=args.fusion_iou_thr,
172+
skip_box_thr=args.skip_box_thr,
173+
conf_type=args.conf_type)
174+
175+
pred_instances = InstanceData()
176+
pred_instances.bboxes = bboxes
177+
pred_instances.scores = scores
178+
pred_instances.labels = labels
179+
180+
fusion_result = DetDataSample(pred_instances=pred_instances)
181+
182+
img_name = filename_list[i]
183+
184+
if not args.no_save_pred:
185+
out_json_path = (
186+
args.out_dir + '/preds/' + img_name.split('.')[0] + '.json')
187+
mmengine.dump(
188+
{
189+
'labels': labels.tolist(),
190+
'scores': scores.tolist(),
191+
'bboxes': bboxes.tolist()
192+
}, out_json_path)
193+
194+
out_file = osp.join(args.out_dir, 'vis',
195+
img_name) if not args.no_save_vis else None
196+
197+
visualizer.add_datasample(
198+
img_name,
199+
inputs[i][..., ::-1],
200+
data_sample=fusion_result,
201+
show=args.show,
202+
draw_gt=False,
203+
wait_time=0,
204+
pred_score_thr=args.pred_score_thr,
205+
out_file=out_file)
206+
207+
if not args.no_save_vis:
208+
print_log(f'results have been saved at {args.out_dir}')
209+
210+
211+
if __name__ == '__main__':
212+
main()

docs/en/user_guides/useful_tools.md

Lines changed: 74 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -111,6 +111,80 @@ python tools/analysis_tools/analyze_results.py \
111111
--show-score-thr 0.3
112112
```
113113

114+
## Fusing results from multiple models
115+
116+
`tools/analysis_tools/fusion_results.py` can fusing predictions using Weighted Boxes Fusion(WBF) from different object detection models. (Currently support coco format only)
117+
118+
**Usage**
119+
120+
```shell
121+
python tools/analysis_tools/fuse_results.py \
122+
${PRED_RESULTS} \
123+
[--annotation ${ANNOTATION}] \
124+
[--weights ${WEIGHTS}] \
125+
[--fusion-iou-thr ${FUSION_IOU_THR}] \
126+
[--skip-box-thr ${SKIP_BOX_THR}] \
127+
[--conf-type ${CONF_TYPE}] \
128+
[--eval-single ${EVAL_SINGLE}] \
129+
[--save-fusion-results ${SAVE_FUSION_RESULTS}] \
130+
[--out-dir ${OUT_DIR}]
131+
```
132+
133+
Description of all arguments:
134+
135+
- `pred-results`: Paths of detection results from different models.(Currently support coco format only)
136+
- `--annotation`: Path of ground-truth.
137+
- `--weights`: List of weights for each model. Default: `None`, which means weight == 1 for each model.
138+
- `--fusion-iou-thr`: IoU value for boxes to be a match。Default: `0.55`
139+
- `--skip-box-thr`: The confidence threshold that needs to be excluded in the WBF algorithm. bboxes whose confidence is less than this value will be excluded.。Default: `0`
140+
- `--conf-type`: How to calculate confidence in weighted boxes.
141+
- `avg`: average value,default.
142+
- `max`: maximum value.
143+
- `box_and_model_avg`: box and model wise hybrid weighted average.
144+
- `absent_model_aware_avg`: weighted average that takes into account the absent model.
145+
- `--eval-single`: Whether evaluate every single model. Default: `False`.
146+
- `--save-fusion-results`: Whether save fusion results. Default: `False`.
147+
- `--out-dir`: Path of fusion results.
148+
149+
**Examples**:
150+
Assume that you have got 3 result files from corresponding models through `tools/test.py`, which paths are './faster-rcnn_r50-caffe_fpn_1x_coco.json', './retinanet_r50-caffe_fpn_1x_coco.json', './cascade-rcnn_r50-caffe_fpn_1x_coco.json' respectively. The ground-truth file path is './annotation.json'.
151+
152+
1. Fusion of predictions from three models and evaluation of their effectiveness
153+
154+
```shell
155+
python tools/analysis_tools/fuse_results.py \
156+
./faster-rcnn_r50-caffe_fpn_1x_coco.json \
157+
./retinanet_r50-caffe_fpn_1x_coco.json \
158+
./cascade-rcnn_r50-caffe_fpn_1x_coco.json \
159+
--annotation ./annotation.json \
160+
--weights 1 2 3 \
161+
```
162+
163+
2. Simultaneously evaluate each single model and fusion results
164+
165+
```shell
166+
python tools/analysis_tools/fuse_results.py \
167+
./faster-rcnn_r50-caffe_fpn_1x_coco.json \
168+
./retinanet_r50-caffe_fpn_1x_coco.json \
169+
./cascade-rcnn_r50-caffe_fpn_1x_coco.json \
170+
--annotation ./annotation.json \
171+
--weights 1 2 3 \
172+
--eval-single
173+
```
174+
175+
3. Fusion of prediction results from three models and save
176+
177+
```shell
178+
python tools/analysis_tools/fuse_results.py \
179+
./faster-rcnn_r50-caffe_fpn_1x_coco.json \
180+
./retinanet_r50-caffe_fpn_1x_coco.json \
181+
./cascade-rcnn_r50-caffe_fpn_1x_coco.json \
182+
--annotation ./annotation.json \
183+
--weights 1 2 3 \
184+
--save-fusion-results \
185+
--out-dir outputs/fusion
186+
```
187+
114188
## Visualization
115189

116190
### Visualize Datasets

docs/zh_cn/user_guides/useful_tools.md

Lines changed: 74 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -109,6 +109,80 @@ python tools/analysis_tools/analyze_results.py \
109109
--show-score-thr 0.3
110110
```
111111

112+
## 多模型检测结果融合
113+
114+
`tools/analysis_tools/fuse_results.py` 可使用 Weighted Boxes Fusion(WBF) 方法将多个模型的检测结果进行融合。(当前仅支持 COCO 格式)
115+
116+
**使用方法**
117+
118+
```shell
119+
python tools/analysis_tools/fuse_results.py \
120+
${PRED_RESULTS} \
121+
[--annotation ${ANNOTATION}] \
122+
[--weights ${WEIGHTS}] \
123+
[--fusion-iou-thr ${FUSION_IOU_THR}] \
124+
[--skip-box-thr ${SKIP_BOX_THR}] \
125+
[--conf-type ${CONF_TYPE}] \
126+
[--eval-single ${EVAL_SINGLE}] \
127+
[--save-fusion-results ${SAVE_FUSION_RESULTS}] \
128+
[--out-dir ${OUT_DIR}]
129+
```
130+
131+
各个参数选项的作用:
132+
133+
- `pred-results`: 多模型测试结果的保存路径。(目前仅支持 json 格式)
134+
- `--annotation`: 真实标注框的保存路径。
135+
- `--weights`: 模型融合权重。默认设置下,每个模型的权重均为1。
136+
- `--fusion-iou-thr`: 在WBF算法中,匹配成功的 IoU 阈值,默认值为`0.55`
137+
- `--skip-box-thr`: WBF算法中需剔除的置信度阈值,置信度小于该值的 bbox 会被剔除,默认值为`0`
138+
- `--conf-type`: 如何计算融合后 bbox 的置信度。有以下四种选项:
139+
- `avg`: 取平均值,默认为此选项。
140+
- `max`: 取最大值。
141+
- `box_and_model_avg`: box和模型尺度的加权平均值。
142+
- `absent_model_aware_avg`: 考虑缺失模型的加权平均值。
143+
- `--eval-single`: 是否评估每个单一模型,默认值为`False`
144+
- `--save-fusion-results`: 是否保存融合结果,默认值为`False`
145+
- `--out-dir`: 融合结果保存的路径。
146+
147+
**样例**:
148+
假设你已经通过 `tools/test.py` 得到了3个模型的 json 格式的结果文件,路径分别为 './faster-rcnn_r50-caffe_fpn_1x_coco.json', './retinanet_r50-caffe_fpn_1x_coco.json', './cascade-rcnn_r50-caffe_fpn_1x_coco.json',真实标注框的文件路径为'./annotation.json'。
149+
150+
1. 融合三个模型的预测结果并评估其效果
151+
152+
```shell
153+
python tools/analysis_tools/fuse_results.py \
154+
./faster-rcnn_r50-caffe_fpn_1x_coco.json \
155+
./retinanet_r50-caffe_fpn_1x_coco.json \
156+
./cascade-rcnn_r50-caffe_fpn_1x_coco.json \
157+
--annotation ./annotation.json \
158+
--weights 1 2 3 \
159+
```
160+
161+
2. 同时评估每个单一模型与融合结果
162+
163+
```shell
164+
python tools/analysis_tools/fuse_results.py \
165+
./faster-rcnn_r50-caffe_fpn_1x_coco.json \
166+
./retinanet_r50-caffe_fpn_1x_coco.json \
167+
./cascade-rcnn_r50-caffe_fpn_1x_coco.json \
168+
--annotation ./annotation.json \
169+
--weights 1 2 3 \
170+
--eval-single
171+
```
172+
173+
3. 融合三个模型的预测结果并保存
174+
175+
```shell
176+
python tools/analysis_tools/fuse_results.py \
177+
./faster-rcnn_r50-caffe_fpn_1x_coco.json \
178+
./retinanet_r50-caffe_fpn_1x_coco.json \
179+
./cascade-rcnn_r50-caffe_fpn_1x_coco.json \
180+
--annotation ./annotation.json \
181+
--weights 1 2 3 \
182+
--save-fusion-results \
183+
--out-dir outputs/fusion
184+
```
185+
112186
## 可视化
113187

114188
### 可视化数据集

mmdet/models/utils/__init__.py

Lines changed: 2 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -19,6 +19,7 @@
1919
from .point_sample import (get_uncertain_point_coords_with_randomness,
2020
get_uncertainty)
2121
from .vlfuse_helper import BertEncoderLayer, VLFuse, permute_and_flatten
22+
from .wbf import weighted_boxes_fusion
2223

2324
__all__ = [
2425
'gaussian_radius', 'gen_gaussian_target', 'make_divisible',
@@ -32,5 +33,5 @@
3233
'samplelist_boxtype2tensor', 'filter_gt_instances', 'rename_loss_dict',
3334
'reweight_loss_dict', 'relative_coordinate_maps', 'aligned_bilinear',
3435
'unfold_wo_center', 'imrenormalize', 'VLFuse', 'permute_and_flatten',
35-
'BertEncoderLayer', 'align_tensor'
36+
'BertEncoderLayer', 'align_tensor', 'weighted_boxes_fusion'
3637
]

0 commit comments

Comments
 (0)