-
Notifications
You must be signed in to change notification settings - Fork 9.6k
Commit
This commit does not belong to any branch on this repository, and may belong to a fork outside of the repository.
Switch imshow_det_bboxes visualization backend from opencv to matplot…
…lib (#4389) * FPN deprecated warning * FPN deprecated warning * Replace imshow_det_bboxes visualization backend * Add bbox_vis unit tests * Encapsulate color_val_matplotlib function * Add fun input parameters * Deprecate block * Add mask display in image * Putting the text inner left corner the bbox * Add a filling color for text regions. * Update color docs * Fix color docs * Update color docs and Fix default param
- Loading branch information
1 parent
9b2c208
commit 308f0d7
Showing
6 changed files
with
290 additions
and
38 deletions.
There are no files selected for viewing
This file contains bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
This file contains bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
Original file line number | Diff line number | Diff line change |
---|---|---|
@@ -0,0 +1,3 @@ | ||
from .image import color_val_matplotlib, imshow_det_bboxes | ||
|
||
__all__ = ['imshow_det_bboxes', 'color_val_matplotlib'] |
This file contains bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
Original file line number | Diff line number | Diff line change |
---|---|---|
@@ -0,0 +1,169 @@ | ||
import os.path as osp | ||
import warnings | ||
|
||
import matplotlib.pyplot as plt | ||
import mmcv | ||
import numpy as np | ||
from matplotlib.collections import PatchCollection | ||
from matplotlib.patches import Polygon | ||
|
||
|
||
def color_val_matplotlib(color): | ||
"""Convert various input in BGR order to normalized RGB matplotlib color | ||
tuples, | ||
Args: | ||
color (:obj:`Color`/str/tuple/int/ndarray): Color inputs | ||
Returns: | ||
tuple[float]: A tuple of 3 normalized floats indicating RGB channels. | ||
""" | ||
color = mmcv.color_val(color) | ||
color = [color / 255 for color in color[::-1]] | ||
return tuple(color) | ||
|
||
|
||
def imshow_det_bboxes(img, | ||
bboxes, | ||
labels, | ||
segms=None, | ||
class_names=None, | ||
score_thr=0, | ||
bbox_color='green', | ||
text_color='green', | ||
mask_color=None, | ||
thickness=2, | ||
font_scale=0.5, | ||
font_size=13, | ||
win_name='', | ||
fig_size=(15, 10), | ||
show=True, | ||
wait_time=0, | ||
out_file=None): | ||
"""Draw bboxes and class labels (with scores) on an image. | ||
Args: | ||
img (str or ndarray): The image to be displayed. | ||
bboxes (ndarray): Bounding boxes (with scores), shaped (n, 4) or | ||
(n, 5). | ||
labels (ndarray): Labels of bboxes. | ||
segms (ndarray or None): Masks, shaped (n,h,w) or None | ||
class_names (list[str]): Names of each classes. | ||
score_thr (float): Minimum score of bboxes to be shown. Default: 0 | ||
bbox_color (str or tuple(int) or :obj:`Color`):Color of bbox lines. | ||
The tuple of color should be in BGR order. Default: 'green' | ||
text_color (str or tuple(int) or :obj:`Color`):Color of texts. | ||
The tuple of color should be in BGR order. Default: 'green' | ||
mask_color (None or str or tuple(int) or :obj:`Color`): | ||
Color of masks. The tuple of color should be in BGR order. | ||
Default: None | ||
thickness (int): Thickness of lines. Default: 2 | ||
font_scale (float): Font scales of texts. Default: 0.5 | ||
font_size (int): Font size of texts. Default: 13 | ||
show (bool): Whether to show the image. Default: True | ||
win_name (str): The window name. Default: '' | ||
fig_size (tuple): Figure size of the pyplot figure. Default: (15, 10) | ||
wait_time (float): Value of waitKey param. Default: 0. | ||
out_file (str or None): The filename to write the image. Default: None | ||
Returns: | ||
ndarray: The image with bboxes drawn on it. | ||
""" | ||
warnings.warn('"font_scale" will be deprecated in v2.9.0,' | ||
'Please use "font_size"') | ||
assert bboxes.ndim == 2, \ | ||
f' bboxes ndim should be 2, but its ndim is {bboxes.ndim}.' | ||
assert labels.ndim == 1, \ | ||
f' labels ndim should be 1, but its ndim is {labels.ndim}.' | ||
assert bboxes.shape[0] == labels.shape[0], \ | ||
'bboxes.shape[0] and labels.shape[0] should have the same length.' | ||
assert bboxes.shape[1] == 4 or bboxes.shape[1] == 5,\ | ||
f' bboxes.shape[1] should be 4 or 5, but its {bboxes.shape[1]}.' | ||
img = mmcv.imread(img).copy() | ||
|
||
if score_thr > 0: | ||
assert bboxes.shape[1] == 5 | ||
scores = bboxes[:, -1] | ||
inds = scores > score_thr | ||
bboxes = bboxes[inds, :] | ||
labels = labels[inds] | ||
if segms is not None: | ||
segms = segms[inds, ...] | ||
|
||
mask_colors = [] | ||
if labels.shape[0] > 0: | ||
if mask_color is None: | ||
# random color | ||
np.random.seed(42) | ||
mask_colors = [ | ||
np.random.randint(0, 256, (1, 3), dtype=np.uint8) | ||
for _ in range(max(labels) + 1) | ||
] | ||
else: | ||
# specify color | ||
mask_colors = [ | ||
np.array(mmcv.color_val(mask_color)[::-1], dtype=np.uint8) | ||
] * ( | ||
max(labels) + 1) | ||
|
||
bbox_color = color_val_matplotlib(bbox_color) | ||
text_color = color_val_matplotlib(text_color) | ||
|
||
img = mmcv.bgr2rgb(img) | ||
img = np.ascontiguousarray(img) | ||
|
||
plt.figure(figsize=fig_size) | ||
plt.title(win_name) | ||
plt.axis('off') | ||
ax = plt.gca() | ||
|
||
polygons = [] | ||
color = [] | ||
for i, (bbox, label) in enumerate(zip(bboxes, labels)): | ||
bbox_int = bbox.astype(np.int32) | ||
poly = [[bbox_int[0], bbox_int[1]], [bbox_int[0], bbox_int[3]], | ||
[bbox_int[2], bbox_int[3]], [bbox_int[2], bbox_int[1]]] | ||
np_poly = np.array(poly).reshape((4, 2)) | ||
polygons.append(Polygon(np_poly)) | ||
color.append(bbox_color) | ||
label_text = class_names[ | ||
label] if class_names is not None else f'class {label}' | ||
if len(bbox) > 4: | ||
label_text += f'|{bbox[-1]:.02f}' | ||
ax.text( | ||
bbox_int[0], | ||
bbox_int[1], | ||
f'{label_text}', | ||
bbox={ | ||
'facecolor': 'black', | ||
'alpha': 0.8, | ||
'pad': 0.7, | ||
'edgecolor': 'none' | ||
}, | ||
color=text_color, | ||
fontsize=font_size, | ||
verticalalignment='top', | ||
horizontalalignment='left') | ||
if segms is not None: | ||
color_mask = mask_colors[labels[i]] | ||
mask = segms[i].astype(bool) | ||
img[mask] = img[mask] * 0.5 + color_mask * 0.5 | ||
|
||
plt.imshow(img) | ||
|
||
p = PatchCollection( | ||
polygons, facecolor='none', edgecolors=color, linewidths=thickness) | ||
ax.add_collection(p) | ||
|
||
if out_file is not None: | ||
dir_name = osp.abspath(osp.dirname(out_file)) | ||
mmcv.mkdir_or_exist(dir_name) | ||
plt.savefig(out_file) | ||
if show: | ||
if wait_time == 0: | ||
plt.show() | ||
else: | ||
plt.show(block=False) | ||
plt.pause(wait_time) | ||
plt.close() | ||
return mmcv.rgb2bgr(img) |
This file contains bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
This file contains bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
Original file line number | Diff line number | Diff line change |
---|---|---|
@@ -0,0 +1,61 @@ | ||
# Copyright (c) Open-MMLab. All rights reserved. | ||
import os.path as osp | ||
import tempfile | ||
|
||
import mmcv | ||
import numpy as np | ||
import pytest | ||
import torch | ||
|
||
from mmdet.core import visualization as vis | ||
|
||
|
||
def test_color(): | ||
assert vis.color_val_matplotlib(mmcv.Color.blue) == (0., 0., 1.) | ||
assert vis.color_val_matplotlib('green') == (0., 1., 0.) | ||
assert vis.color_val_matplotlib((1, 2, 3)) == (3 / 255, 2 / 255, 1 / 255) | ||
assert vis.color_val_matplotlib(100) == (100 / 255, 100 / 255, 100 / 255) | ||
assert vis.color_val_matplotlib(np.zeros(3, dtype=np.int)) == (0., 0., 0.) | ||
# forbid white color | ||
with pytest.raises(TypeError): | ||
vis.color_val_matplotlib([255, 255, 255]) | ||
# forbid float | ||
with pytest.raises(TypeError): | ||
vis.color_val_matplotlib(1.0) | ||
# overflowed | ||
with pytest.raises(AssertionError): | ||
vis.color_val_matplotlib((0, 0, 500)) | ||
|
||
|
||
def test_imshow_det_bboxes(): | ||
tmp_filename = osp.join(tempfile.gettempdir(), 'det_bboxes_image', | ||
'image.jpg') | ||
image = np.ones((10, 10, 3), np.uint8) | ||
bbox = np.array([[2, 1, 3, 3], [3, 4, 6, 6]]) | ||
label = np.array([0, 1]) | ||
vis.imshow_det_bboxes( | ||
image, bbox, label, out_file=tmp_filename, show=False) | ||
assert osp.isfile(tmp_filename) | ||
|
||
# test shaped (0,) | ||
image = np.ones((10, 10, 3), np.uint8) | ||
bbox = np.ones((0, 4)) | ||
label = np.ones((0, )) | ||
vis.imshow_det_bboxes( | ||
image, bbox, label, out_file=tmp_filename, show=False) | ||
|
||
# test mask | ||
image = np.ones((10, 10, 3), np.uint8) | ||
bbox = np.array([[2, 1, 3, 3], [3, 4, 6, 6]]) | ||
label = np.array([0, 1]) | ||
segms = np.random.random((2, 10, 10)) > 0.5 | ||
segms = np.array(segms, np.int32) | ||
vis.imshow_det_bboxes( | ||
image, bbox, label, segms, out_file=tmp_filename, show=False) | ||
assert osp.isfile(tmp_filename) | ||
|
||
# test tensor mask type error | ||
with pytest.raises(AttributeError): | ||
segms = torch.tensor(segms) | ||
vis.imshow_det_bboxes( | ||
image, bbox, label, segms, out_file=tmp_filename, show=False) |
Oops, something went wrong.