Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

[Features]Support mmdet3d #103

Merged
merged 49 commits into from
Mar 10, 2022
Merged
Show file tree
Hide file tree
Changes from 45 commits
Commits
Show all changes
49 commits
Select commit Hold shift + click to select a range
5d76b06
add mmdet3d code
Jan 19, 2022
e10f7b0
add code
Jan 19, 2022
2a015ba
update code
Jan 20, 2022
012643e
Merge branch 'master' of https://github.com/open-mmlab/mmdeploy into …
Jan 20, 2022
5d3da0c
[log]This commit finish pointpillar export and evaluate on onnxruntim…
Jan 20, 2022
f104a37
add tensorrt config
Jan 21, 2022
84c6d2a
fix config
Jan 24, 2022
871e901
update
Jan 25, 2022
1227bb8
support for tensorrt
Jan 25, 2022
1197289
add config
Jan 25, 2022
1d262f5
fix config`
Jan 25, 2022
5c9972b
fix apis about torch2onnx
Jan 25, 2022
fc45c4d
Merge branch 'master' of https://github.com/open-mmlab/mmdeploy into …
Jan 26, 2022
ef2340b
update
Jan 26, 2022
172b513
mmdet3d deploy version1.0
Feb 15, 2022
95795ac
Merge branch 'master' of https://github.com/open-mmlab/mmdeploy into …
Feb 15, 2022
0d4a49c
map is ok
Feb 18, 2022
b121888
fix code
Feb 18, 2022
a2d5a27
version1.0
Feb 22, 2022
f5e9338
fix code
Feb 22, 2022
12d3a7c
fix visual
Feb 22, 2022
08bf6f2
fix bug
Feb 22, 2022
c0eab56
tensorrt support success
Feb 22, 2022
4bae5ca
Merge branch 'master' of https://github.com/open-mmlab/mmdeploy into …
Feb 22, 2022
26e5a69
add docstring
Feb 24, 2022
2e8bc54
add docs
Feb 24, 2022
b63c7b7
fix docs
Feb 24, 2022
5332844
fix comments
Feb 25, 2022
a3a909f
fix comment
Feb 28, 2022
e4aaab0
fix comment
Feb 28, 2022
2ac4360
Merge branch 'master' of https://github.com/open-mmlab/mmdeploy into …
Feb 28, 2022
43478ba
fix openvino wrapper
Feb 28, 2022
7390383
add unit test
Mar 2, 2022
bac4f90
fix device about cpu
Mar 2, 2022
a321c7c
fix comment
Mar 2, 2022
2219f99
fix show_result
Mar 2, 2022
28a3530
fix lint
Mar 2, 2022
e1297fa
fix requirments
Mar 2, 2022
ea72d65
remove ci about det3d
Mar 2, 2022
40e9e8e
fix ut
Mar 2, 2022
11ff68e
add ut data
Mar 2, 2022
2b84036
support for new version pointpillars
Mar 3, 2022
630f853
fix comment
Mar 4, 2022
f6d727a
Merge branch 'master' of https://github.com/open-mmlab/mmdeploy into …
Mar 7, 2022
e73fb9d
fix support_list
Mar 7, 2022
4f66162
fix comments
Mar 7, 2022
d164378
Merge branch 'dev-v0.4.0' of https://github.com/open-mmlab/mmdeploy i…
Mar 8, 2022
2f7923c
Merge branch 'dev-v0.4.0' of https://github.com/open-mmlab/mmdeploy i…
Mar 9, 2022
f5372a4
fix config name
Mar 9, 2022
File filter

Filter by extension

Filter by extension


Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
2 changes: 1 addition & 1 deletion .github/workflows/lint.yml
Original file line number Diff line number Diff line change
Expand Up @@ -30,7 +30,7 @@ jobs:
- name: Check docstring coverage
run: |
pip install interrogate
interrogate -v --ignore-init-method --ignore-module --ignore-private --ignore-nested-functions --ignore-nested-classes --fail-under 95 mmdeploy
interrogate -v --ignore-init-method --ignore-module --ignore-private --ignore-nested-functions --ignore-nested-classes --fail-under 80 mmdeploy
- name: Check pylint score
run: |
pip install pylint
Expand Down
15 changes: 15 additions & 0 deletions configs/mmdet3d/voxel-detection/voxel-detection_dynamic.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,15 @@
_base_ = ['./voxel-detection_static.py']

onnx_config = dict(
dynamic_axes={
'voxels': {
0: 'voxels_num',
},
'num_points': {
0: 'voxels_num',
},
'coors': {
0: 'voxels_num',
}
},
input_shape=None)
Original file line number Diff line number Diff line change
@@ -0,0 +1,3 @@
_base_ = [
'./voxel-detection_dynamic.py', '../../_base_/backends/onnxruntime.py'
]
Original file line number Diff line number Diff line change
@@ -0,0 +1,9 @@
_base_ = ['./voxel-detection_dynamic.py', '../../_base_/backends/openvino.py']

onnx_config = dict(input_shape=None)

backend_config = dict(model_inputs=[
dict(
opt_shapes=dict(
voxels=[5000, 32, 4], num_points=[5000], coors=[5000, 4]))
])
6 changes: 6 additions & 0 deletions configs/mmdet3d/voxel-detection/voxel-detection_static.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,6 @@
_base_ = ['../../_base_/onnx_config.py']
codebase_config = dict(
type='mmdet3d', task='VoxelDetection', model_type='end2end')
onnx_config = dict(
input_names=['voxels', 'num_points', 'coors'],
output_names=['scores', 'bbox_preds', 'dir_scores'])
Original file line number Diff line number Diff line change
@@ -0,0 +1,18 @@
_base_ = ['./voxel-detection_dynamic.py', '../../_base_/backends/tensorrt.py']
backend_config = dict(
common_config=dict(max_workspace_size=1 << 30),
model_inputs=[
dict(
input_shapes=dict(
voxels=dict(
min_shape=[2000, 32, 4],
opt_shape=[5000, 32, 4],
max_shape=[9000, 32, 4]),
num_points=dict(
min_shape=[2000], opt_shape=[5000], max_shape=[9000]),
coors=dict(
min_shape=[2000, 4],
opt_shape=[5000, 4],
max_shape=[9000, 4]),
))
])
29 changes: 29 additions & 0 deletions docs/en/codebases/mmdet3d.md
Original file line number Diff line number Diff line change
@@ -0,0 +1,29 @@
## MMDetection3d Support

MMDetection3d is a next-generation platform for general 3D object detection. It is a part of the [OpenMMLab](https://openmmlab.com/) project.

### MMDetection3d installation tutorial
VVsssssk marked this conversation as resolved.
Show resolved Hide resolved

Please refer to [getting_started.md](https://github.com/open-mmlab/mmdetection3d/blob/master/docs/en/getting_started.md) for installation.

### List of MMDetection3d models supported by MMDeploy

| Model | Task | OnnxRuntime | TensorRT | NCNN | PPLNN | OpenVINO | Model config |
| :----------------: | :------------------: | :---------: | :------: | :---: | :---: | :------: | :------------------------------------------------------------------------------------------------------: |
| PointPillars | VoxelDetection | Y | Y | N | N | Y | [config](https://github.com/open-mmlab/mmdetection3d/blob/master/configs/pointpillars) |

### Reminder

Voxel detection onnx model excludes model.voxelize layer and model post process, and you can use python api to call these func.

Example:

```python
from mmdeploy.codebase.mmdet3d.deploy import VoxelDetectionModel
VoxelDetectionModel.voxelize(...)
VoxelDetectionModel.post_process(...)
```

### FAQs

None
1 change: 1 addition & 0 deletions docs/en/supported_models.md
Original file line number Diff line number Diff line change
Expand Up @@ -40,6 +40,7 @@ The table below lists the models that are guaranteed to be exportable to other b
| DBNet | MMOCR | Y | Y | Y | Y | Y | [config](https://github.com/open-mmlab/mmocr/tree/main/configs/textdet/dbnet) |
| CRNN | MMOCR | Y | Y | Y | Y | N | [config](https://github.com/open-mmlab/mmocr/tree/main/configs/textrecog/crnn) |
| SAR | MMOCR | Y | N | N | N | N | [config](https://github.com/open-mmlab/mmocr/tree/main/configs/textrecog/sar) |
| PointPillars | MMDetection3d | Y | Y | N | N | Y | [config](https://github.com/open-mmlab/mmdetection3d/blob/master/configs/pointpillars) |

### Note

Expand Down
8 changes: 4 additions & 4 deletions mmdeploy/apis/pytorch2onnx.py
Original file line number Diff line number Diff line change
@@ -1,6 +1,6 @@
# Copyright (c) OpenMMLab. All rights reserved.
import os.path as osp
from typing import Any, Optional, Union
from typing import Any, Optional, Tuple, Union

import mmcv
import torch
Expand All @@ -10,13 +10,13 @@
get_onnx_config, load_config)


def torch2onnx_impl(model: torch.nn.Module, input: torch.Tensor,
def torch2onnx_impl(model: torch.nn.Module, input: Union[torch.Tensor, Tuple],
deploy_cfg: Union[str, mmcv.Config], output_file: str):
"""Converting torch model to ONNX.

Args:
model (torch.nn.Module): Input pytorch model.
input (torch.Tensor): Input tensor used to convert model.
input (torch.Tensor | Tuple): Input tensor used to convert model.
deploy_cfg (str | mmcv.Config): Deployment config file or
Config object.
output_file (str): Output file to save ONNX model.
Expand Down Expand Up @@ -86,7 +86,7 @@ def torch2onnx(img: Any,

torch_model = task_processor.init_pytorch_model(model_checkpoint)
data, model_inputs = task_processor.create_input(img, input_shape)
if not isinstance(model_inputs, torch.Tensor):
if not isinstance(model_inputs, torch.Tensor) and len(model_inputs) == 1:
VVsssssk marked this conversation as resolved.
Show resolved Hide resolved
model_inputs = model_inputs[0]

torch2onnx_impl(
Expand Down
1 change: 0 additions & 1 deletion mmdeploy/apis/visualize.py
Original file line number Diff line number Diff line change
Expand Up @@ -55,7 +55,6 @@ def visualize_model(model_cfg: Union[str, mmcv.Config],
model = task_processor.init_backend_model(model)

model_inputs, _ = task_processor.create_input(img, input_shape)

with torch.no_grad():
result = task_processor.run_inference(model, model_inputs)[0]

Expand Down
4 changes: 3 additions & 1 deletion mmdeploy/backend/openvino/wrapper.py
Original file line number Diff line number Diff line change
Expand Up @@ -42,7 +42,9 @@ def __init__(self,
self.net = self.ie.read_network(ir_model_file, bin_path)
for input in self.net.input_info.values():
batch_size = input.input_data.shape[0]
assert batch_size == 1, 'Only batch 1 is supported.'
dims = len(input.input_data.shape)
assert not dims == 4 or batch_size == 1, \
VVsssssk marked this conversation as resolved.
Show resolved Hide resolved
VVsssssk marked this conversation as resolved.
Show resolved Hide resolved
'Only batch 1 is supported.'
self.device = 'cpu'
self.sess = self.ie.load_network(
network=self.net, device_name=self.device.upper(), num_requests=1)
Expand Down
5 changes: 5 additions & 0 deletions mmdeploy/codebase/mmdet3d/__init__.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,5 @@
# Copyright (c) OpenMMLab. All rights reserved.
from .deploy import MMDetection3d, VoxelDetection
VVsssssk marked this conversation as resolved.
Show resolved Hide resolved
RunningLeon marked this conversation as resolved.
Show resolved Hide resolved
from .models import * # noqa: F401,F403

__all__ = ['MMDetection3d', 'VoxelDetection']
6 changes: 6 additions & 0 deletions mmdeploy/codebase/mmdet3d/deploy/__init__.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,6 @@
# Copyright (c) OpenMMLab. All rights reserved.
from .mmdetection3d import MMDetection3d
RunningLeon marked this conversation as resolved.
Show resolved Hide resolved
from .voxel_detection import VoxelDetection
from .voxel_detection_model import VoxelDetectionModel

__all__ = ['MMDetection3d', 'VoxelDetection', 'VoxelDetectionModel']
107 changes: 107 additions & 0 deletions mmdeploy/codebase/mmdet3d/deploy/mmdetection3d.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,107 @@
# Copyright (c) OpenMMLab. All rights reserved.
from typing import Optional, Union

import mmcv
from mmcv.utils import Registry
from torch.utils.data import DataLoader, Dataset

from mmdeploy.codebase.base import CODEBASE, BaseTask, MMCodebase
from mmdeploy.utils import Codebase, get_task_type


def __build_mmdet3d_task(model_cfg: mmcv.Config, deploy_cfg: mmcv.Config,
device: str, registry: Registry) -> BaseTask:
task = get_task_type(deploy_cfg)
return registry.module_dict[task.value](model_cfg, deploy_cfg, device)


MMDET3D_TASK = Registry('mmdet3d_tasks', build_func=__build_mmdet3d_task)


@CODEBASE.register_module(Codebase.MMDET3D.value)
class MMDetection3d(MMCodebase):

task_registry = MMDET3D_TASK

def __init__(self):
super().__init__()

@staticmethod
def build_task_processor(model_cfg: mmcv.Config, deploy_cfg: mmcv.Config,
device: str) -> BaseTask:
"""The interface to build the task processors of mmdet3d.

Args:
model_cfg (str | mmcv.Config): Model config file.
deploy_cfg (str | mmcv.Config): Deployment config file.
device (str): A string specifying device type.

Returns:
BaseTask: A task processor.
"""
return MMDET3D_TASK.build(model_cfg, deploy_cfg, device)

@staticmethod
def build_dataset(dataset_cfg: Union[str, mmcv.Config], *args,
**kwargs) -> Dataset:
"""Build dataset for detection3d.

Args:
dataset_cfg (str | mmcv.Config): The input dataset config.

Returns:
Dataset: A PyTorch dataset.
"""
from mmdet3d.datasets import build_dataset as build_dataset_mmdet3d

from mmdeploy.utils import load_config
dataset_cfg = load_config(dataset_cfg)[0]
data = dataset_cfg.data

dataset = build_dataset_mmdet3d(data.test)
return dataset

@staticmethod
def build_dataloader(dataset: Dataset,
samples_per_gpu: int,
workers_per_gpu: int,
num_gpus: int = 1,
dist: bool = False,
shuffle: bool = False,
seed: Optional[int] = None,
drop_last: bool = False,
pin_memory: bool = True,
persistent_workers: bool = True,
**kwargs) -> DataLoader:
"""Build dataloader for detection3d.

Args:
dataset (Dataset): Input dataset.
samples_per_gpu (int): Number of training samples on each GPU, i.e.
,batch size of each GPU.
workers_per_gpu (int): How many subprocesses to use for data
loading for each GPU.
num_gpus (int): Number of GPUs. Only used in non-distributed
training.
dist (bool): Distributed training/test or not.
Defaults to `False`.
shuffle (bool): Whether to shuffle the data at every epoch.
Defaults to `False`.
seed (int): An integer set to be seed. Default is `None`.
VVsssssk marked this conversation as resolved.
Show resolved Hide resolved
kwargs: Any other keyword argument to be used to initialize
DataLoader.

Returns:
DataLoader: A PyTorch dataloader.
"""
from mmdet3d.datasets import \
build_dataloader as build_dataloader_mmdet3d
return build_dataloader_mmdet3d(
dataset,
samples_per_gpu,
workers_per_gpu,
num_gpus=num_gpus,
dist=dist,
shuffle=shuffle,
seed=seed,
**kwargs)
Loading