1
1
# Copyright (c) OpenMMLab. All rights reserved.
2
2
import math
3
3
from numbers import Number
4
- from typing import Dict , List , Optional , Sequence , Tuple , Union
4
+ from typing import Dict , List , Optional , Sequence , Union
5
5
6
6
import numpy as np
7
7
import torch
@@ -28,24 +28,25 @@ class Det3DDataPreprocessor(DetDataPreprocessor):
28
28
- 1) For image data:
29
29
- Pad images in inputs to the maximum size of current batch with defined
30
30
``pad_value``. The padding size can be divisible by a defined
31
- ``pad_size_divisor``
31
+ ``pad_size_divisor``.
32
32
- Stack images in inputs to batch_imgs.
33
33
- Convert images in inputs from bgr to rgb if the shape of input is
34
- (3, H, W).
34
+ (3, H, W).
35
35
- Normalize images in inputs with defined std and mean.
36
36
- Do batch augmentations during training.
37
37
38
38
- 2) For point cloud data:
39
- - if no voxelization, directly return list of point cloud data.
40
- - if voxelization is applied, voxelize point cloud according to
39
+ - If no voxelization, directly return list of point cloud data.
40
+ - If voxelization is applied, voxelize point cloud according to
41
41
``voxel_type`` and obtain ``voxels``.
42
42
43
43
Args:
44
- voxel (bool): Whether to apply voxelziation to point cloud.
44
+ voxel (bool): Whether to apply voxelization to point cloud.
45
+ Defaults to False.
45
46
voxel_type (str): Voxelization type. Two voxelization types are
46
47
provided: 'hard' and 'dynamic', respectively for hard
47
48
voxelization and dynamic voxelization. Defaults to 'hard'.
48
- voxel_layer (:obj:`ConfigDict`, optional): Voxelization layer
49
+ voxel_layer (dict or :obj:`ConfigDict`, optional): Voxelization layer
49
50
config. Defaults to None.
50
51
mean (Sequence[Number], optional): The pixel mean of R, G, B channels.
51
52
Defaults to None.
@@ -54,11 +55,21 @@ class Det3DDataPreprocessor(DetDataPreprocessor):
54
55
pad_size_divisor (int): The size of padded image should be
55
56
divisible by ``pad_size_divisor``. Defaults to 1.
56
57
pad_value (Number): The padded pixel value. Defaults to 0.
57
- bgr_to_rgb (bool): whether to convert image from BGR to RGB.
58
+ pad_mask (bool): Whether to pad instance masks. Defaults to False.
59
+ mask_pad_value (int): The padded pixel value for instance masks.
60
+ Defaults to 0.
61
+ pad_seg (bool): Whether to pad semantic segmentation maps.
62
+ Defaults to False.
63
+ seg_pad_value (int): The padded pixel value for semantic
64
+ segmentation maps. Defaults to 255.
65
+ bgr_to_rgb (bool): Whether to convert image from BGR to RGB.
58
66
Defaults to False.
59
- rgb_to_bgr (bool): whether to convert image from RGB to RGB .
67
+ rgb_to_bgr (bool): Whether to convert image from RGB to BGR .
60
68
Defaults to False.
61
- batch_augments (list[dict], optional): Batch-level augmentations
69
+ boxtype2tensor (bool): Whether to keep the ``BaseBoxes`` type of
70
+ bboxes data or not. Defaults to True.
71
+ batch_augments (List[dict], optional): Batch-level augmentations.
72
+ Defaults to None.
62
73
"""
63
74
64
75
def __init__ (self ,
@@ -76,8 +87,8 @@ def __init__(self,
76
87
bgr_to_rgb : bool = False ,
77
88
rgb_to_bgr : bool = False ,
78
89
boxtype2tensor : bool = True ,
79
- batch_augments : Optional [List [dict ]] = None ):
80
- super ().__init__ (
90
+ batch_augments : Optional [List [dict ]] = None ) -> None :
91
+ super (Det3DDataPreprocessor ).__init__ (
81
92
mean = mean ,
82
93
std = std ,
83
94
pad_size_divisor = pad_size_divisor ,
@@ -94,24 +105,21 @@ def __init__(self,
94
105
if voxel :
95
106
self .voxel_layer = Voxelization (** voxel_layer )
96
107
97
- def forward (
98
- self ,
99
- data : Union [dict , List [dict ]],
100
- training : bool = False
101
- ) -> Tuple [Union [dict , List [dict ]], Optional [list ]]:
102
- """Perform normalization、padding and bgr2rgb conversion based on
108
+ def forward (self ,
109
+ data : Union [dict , List [dict ]],
110
+ training : bool = False ) -> Union [dict , List [dict ]]:
111
+ """Perform normalization, padding and bgr2rgb conversion based on
103
112
``BaseDataPreprocessor``.
104
113
105
114
Args:
106
- data (dict | List[dict]): data from dataloader.
115
+ data (dict or List[dict]): Data from dataloader.
107
116
The dict contains the whole batch data, when it is
108
117
a list[dict], the list indicate test time augmentation.
109
-
110
118
training (bool): Whether to enable training time augmentation.
111
119
Defaults to False.
112
120
113
121
Returns:
114
- Dict | List[Dict ]: Data in the same format as the model input.
122
+ dict or List[dict ]: Data in the same format as the model input.
115
123
"""
116
124
if isinstance (data , list ):
117
125
num_augs = len (data )
@@ -126,7 +134,7 @@ def forward(
126
134
return self .simple_process (data , training )
127
135
128
136
def simple_process (self , data : dict , training : bool = False ) -> dict :
129
- """Perform normalization、 padding and bgr2rgb conversion for img data
137
+ """Perform normalization, padding and bgr2rgb conversion for img data
130
138
based on ``BaseDataPreprocessor``, and voxelize point cloud if `voxel`
131
139
is set to be True.
132
140
@@ -188,7 +196,7 @@ def simple_process(self, data: dict, training: bool = False) -> dict:
188
196
189
197
return {'inputs' : batch_inputs , 'data_samples' : data_samples }
190
198
191
- def preprocess_img (self , _batch_img ) :
199
+ def preprocess_img (self , _batch_img : torch . Tensor ) -> torch . Tensor :
192
200
# channel transform
193
201
if self ._channel_conversion :
194
202
_batch_img = _batch_img [[2 , 1 , 0 ], ...]
@@ -206,7 +214,7 @@ def preprocess_img(self, _batch_img):
206
214
return _batch_img
207
215
208
216
def collate_data (self , data : dict ) -> dict :
209
- """Copying data to the target device and Performs normalization、
217
+ """Copying data to the target device and Performs normalization,
210
218
padding and bgr2rgb conversion and stack based on
211
219
``BaseDataPreprocessor``.
212
220
@@ -273,7 +281,7 @@ def collate_data(self, data: dict) -> dict:
273
281
raise TypeError (
274
282
'Output of `cast_data` should be a list of dict '
275
283
'or a tuple with inputs and data_samples, but got'
276
- f'{ type (data )} : { data } ' )
284
+ f'{ type (data )} : { data } ' )
277
285
278
286
data ['inputs' ]['imgs' ] = batch_imgs
279
287
@@ -284,14 +292,14 @@ def collate_data(self, data: dict) -> dict:
284
292
def _get_pad_shape (self , data : dict ) -> List [tuple ]:
285
293
"""Get the pad_shape of each image based on data and
286
294
pad_size_divisor."""
287
- # rewrite `_get_pad_shape` for obaining image inputs.
295
+ # rewrite `_get_pad_shape` for obtaining image inputs.
288
296
_batch_inputs = data ['inputs' ]['img' ]
289
297
# Process data with `pseudo_collate`.
290
298
if is_list_of (_batch_inputs , torch .Tensor ):
291
299
batch_pad_shape = []
292
300
for ori_input in _batch_inputs :
293
301
if ori_input .dim () == 4 :
294
- # mean multiivew input, select ont of the
302
+ # mean multiview input, select one of the
295
303
# image to calculate the pad shape
296
304
ori_input = ori_input [0 ]
297
305
pad_h = int (
@@ -316,24 +324,24 @@ def _get_pad_shape(self, data: dict) -> List[tuple]:
316
324
batch_pad_shape = [(pad_h , pad_w )] * _batch_inputs .shape [0 ]
317
325
else :
318
326
raise TypeError ('Output of `cast_data` should be a list of dict '
319
- 'or a tuple with inputs and data_samples, but got'
327
+ 'or a tuple with inputs and data_samples, but got '
320
328
f'{ type (data )} : { data } ' )
321
329
return batch_pad_shape
322
330
323
331
@torch .no_grad ()
324
- def voxelize (self , points : List [torch .Tensor ]) -> Dict :
332
+ def voxelize (self , points : List [torch .Tensor ]) -> Dict [ str , torch . Tensor ] :
325
333
"""Apply voxelization to point cloud.
326
334
327
335
Args:
328
336
points (List[Tensor]): Point cloud in one data batch.
329
337
330
338
Returns:
331
- dict [str, Tensor]: Voxelization information.
339
+ Dict [str, Tensor]: Voxelization information.
332
340
333
- - voxels (Tensor): Features of voxels, shape is MXNxC for hard
334
- voxelization, NXC for dynamic voxelization.
335
- - coors (Tensor): Coordinates of voxels, shape is Nx(1+NDim),
336
- where 1 represents the batch index.
341
+ - voxels (Tensor): Features of voxels, shape is MxNxC for hard
342
+ voxelization, NxC for dynamic voxelization.
343
+ - coors (Tensor): Coordinates of voxels, shape is Nx(1+NDim),
344
+ where 1 represents the batch index.
337
345
- num_points (Tensor, optional): Number of points in each voxel.
338
346
- voxel_centers (Tensor, optional): Centers of voxels.
339
347
"""
@@ -342,43 +350,38 @@ def voxelize(self, points: List[torch.Tensor]) -> Dict:
342
350
343
351
if self .voxel_type == 'hard' :
344
352
voxels , coors , num_points , voxel_centers = [], [], [], []
345
- for res in points :
353
+ for i , res in enumerate ( points ) :
346
354
res_voxels , res_coors , res_num_points = self .voxel_layer (res )
347
355
res_voxel_centers = (
348
356
res_coors [:, [2 , 1 , 0 ]] + 0.5 ) * res_voxels .new_tensor (
349
357
self .voxel_layer .voxel_size ) + res_voxels .new_tensor (
350
358
self .voxel_layer .point_cloud_range [0 :3 ])
359
+ res_coors = F .pad (res_coors , (1 , 0 ), mode = 'constant' , value = i )
351
360
voxels .append (res_voxels )
352
361
coors .append (res_coors )
353
362
num_points .append (res_num_points )
354
363
voxel_centers .append (res_voxel_centers )
355
364
356
365
voxels = torch .cat (voxels , dim = 0 )
366
+ coors = torch .cat (coors , dim = 0 )
357
367
num_points = torch .cat (num_points , dim = 0 )
358
368
voxel_centers = torch .cat (voxel_centers , dim = 0 )
359
- coors_batch = []
360
- for i , coor in enumerate (coors ):
361
- coor_pad = F .pad (coor , (1 , 0 ), mode = 'constant' , value = i )
362
- coors_batch .append (coor_pad )
363
- coors_batch = torch .cat (coors_batch , dim = 0 )
369
+
364
370
voxel_dict ['num_points' ] = num_points
365
371
voxel_dict ['voxel_centers' ] = voxel_centers
366
372
elif self .voxel_type == 'dynamic' :
367
373
coors = []
368
374
# dynamic voxelization only provide a coors mapping
369
- for res in points :
375
+ for i , res in enumerate ( points ) :
370
376
res_coors = self .voxel_layer (res )
377
+ res_coors = F .pad (res_coors , (1 , 0 ), mode = 'constant' , value = i )
371
378
coors .append (res_coors )
372
379
voxels = torch .cat (points , dim = 0 )
373
- coors_batch = []
374
- for i , coor in enumerate (coors ):
375
- coor_pad = F .pad (coor , (1 , 0 ), mode = 'constant' , value = i )
376
- coors_batch .append (coor_pad )
377
- coors_batch = torch .cat (coors_batch , dim = 0 )
380
+ coors = torch .cat (coors , dim = 0 )
378
381
else :
379
382
raise ValueError (f'Invalid voxelization type { self .voxel_type } ' )
380
383
381
384
voxel_dict ['voxels' ] = voxels
382
- voxel_dict ['coors' ] = coors_batch
385
+ voxel_dict ['coors' ] = coors
383
386
384
387
return voxel_dict
0 commit comments