Skip to content

Commit ac16582

Browse files
authored
Merge pull request #32 from roboflow/feature/detections_api_refinement
feature/detections_api_refinement
2 parents 3ea95b3 + 8e72283 commit ac16582

File tree

4 files changed

+111
-13
lines changed

4 files changed

+111
-13
lines changed

README.md

Lines changed: 5 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -7,6 +7,7 @@
77
>
88
</a>
99
</p>
10+
1011
<br>
1112

1213
<div align="center">
@@ -53,6 +54,10 @@
5354
</a>
5455
</div>
5556

57+
<br>
58+
59+
[![Colab](https://colab.research.google.com/assets/colab-badge.svg)](https://colab.research.google.com/github/roboflow-ai/notebooks/blob/main/notebooks/how-to-detect-and-count-objects-in-polygon-zone.ipynb)
60+
5661
</div>
5762

5863
## 👋 hello

supervision/__init__.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -1,8 +1,8 @@
1-
__version__ = "0.2.1"
1+
__version__ = "0.3.0"
22

33
from supervision.detection.core import BoxAnnotator, Detections
4-
from supervision.detection.polygon_zone import PolygonZone, PolygonZoneAnnotator
54
from supervision.detection.line_counter import LineZone, LineZoneAnnotator
5+
from supervision.detection.polygon_zone import PolygonZone, PolygonZoneAnnotator
66
from supervision.detection.utils import generate_2d_mask
77
from supervision.draw.color import Color, ColorPalette
88
from supervision.draw.utils import draw_filled_rectangle, draw_polygon, draw_text

supervision/detection/core.py

Lines changed: 67 additions & 11 deletions
Original file line numberDiff line numberDiff line change
@@ -1,11 +1,12 @@
11
from __future__ import annotations
22

33
from dataclasses import dataclass
4-
from typing import List, Optional, Union
4+
from typing import Iterator, List, Optional, Tuple, Union
55

66
import cv2
77
import numpy as np
88

9+
from supervision.detection.utils import non_max_suppression
910
from supervision.draw.color import Color, ColorPalette
1011
from supervision.geometry.core import Position
1112

@@ -17,22 +18,26 @@ class Detections:
1718
1819
Attributes:
1920
xyxy (np.ndarray): An array of shape `(n, 4)` containing the bounding boxes coordinates in format `[x1, y1, x2, y2]`
20-
confidence (np.ndarray): An array of shape `(n,)` containing the confidence scores of the detections.
21+
confidence (Optional[np.ndarray]): An array of shape `(n,)` containing the confidence scores of the detections.
2122
class_id (np.ndarray): An array of shape `(n,)` containing the class ids of the detections.
2223
tracker_id (Optional[np.ndarray]): An array of shape `(n,)` containing the tracker ids of the detections.
2324
"""
2425

2526
xyxy: np.ndarray
26-
confidence: np.ndarray
2727
class_id: np.ndarray
28+
confidence: Optional[np.ndarray] = None
2829
tracker_id: Optional[np.ndarray] = None
2930

3031
def __post_init__(self):
3132
n = len(self.xyxy)
3233
validators = [
3334
(isinstance(self.xyxy, np.ndarray) and self.xyxy.shape == (n, 4)),
34-
(isinstance(self.confidence, np.ndarray) and self.confidence.shape == (n,)),
3535
(isinstance(self.class_id, np.ndarray) and self.class_id.shape == (n,)),
36+
self.confidence is None
37+
or (
38+
isinstance(self.confidence, np.ndarray)
39+
and self.confidence.shape == (n,)
40+
),
3641
self.tracker_id is None
3742
or (
3843
isinstance(self.tracker_id, np.ndarray)
@@ -42,7 +47,7 @@ def __post_init__(self):
4247
if not all(validators):
4348
raise ValueError(
4449
"xyxy must be 2d np.ndarray with (n, 4) shape, "
45-
"confidence must be 1d np.ndarray with (n,) shape, "
50+
"confidence must be None or 1d np.ndarray with (n,) shape, "
4651
"class_id must be 1d np.ndarray with (n,) shape, "
4752
"tracker_id must be None or 1d np.ndarray with (n,) shape"
4853
)
@@ -53,14 +58,16 @@ def __len__(self):
5358
"""
5459
return len(self.xyxy)
5560

56-
def __iter__(self):
61+
def __iter__(
62+
self,
63+
) -> Iterator[Tuple[np.ndarray, Optional[float], int, Optional[Union[str, int]]]]:
5764
"""
5865
Iterates over the Detections object and yield a tuple of `(xyxy, confidence, class_id, tracker_id)` for each detection.
5966
"""
6067
for i in range(len(self.xyxy)):
6168
yield (
6269
self.xyxy[i],
63-
self.confidence[i],
70+
self.confidence[i] if self.confidence is not None else None,
6471
self.class_id[i],
6572
self.tracker_id[i] if self.tracker_id is not None else None,
6673
)
@@ -69,11 +76,17 @@ def __eq__(self, other: Detections):
6976
return all(
7077
[
7178
np.array_equal(self.xyxy, other.xyxy),
72-
np.array_equal(self.confidence, other.confidence),
79+
any(
80+
[
81+
self.confidence is None and other.confidence is None,
82+
np.array_equal(self.confidence, other.confidence),
83+
]
84+
),
7385
np.array_equal(self.class_id, other.class_id),
7486
any(
7587
[
7688
self.tracker_id is None and other.tracker_id is None,
89+
np.array_equal(self.tracker_id, other.tracker_id),
7790
]
7891
),
7992
]
@@ -122,7 +135,7 @@ def from_yolov8(cls, yolov8_results):
122135
>>> from supervision import Detections
123136
124137
>>> model = YOLO('yolov8s.pt')
125-
>>> results = model(frame)
138+
>>> results = model(frame)[0]
126139
>>> detections = Detections.from_yolov8(results)
127140
```
128141
"""
@@ -132,6 +145,36 @@ def from_yolov8(cls, yolov8_results):
132145
class_id=yolov8_results.boxes.cls.cpu().numpy().astype(int),
133146
)
134147

148+
@classmethod
149+
def from_transformers(cls, transformers_results: dict):
150+
return cls(
151+
xyxy=transformers_results["boxes"].cpu().numpy(),
152+
confidence=transformers_results["scores"].cpu().numpy(),
153+
class_id=transformers_results["labels"].cpu().numpy().astype(int),
154+
)
155+
156+
@classmethod
157+
def from_detectron2(cls, detectron2_results):
158+
return cls(
159+
xyxy=detectron2_results["instances"].pred_boxes.tensor.cpu().numpy(),
160+
confidence=detectron2_results["instances"].scores.cpu().numpy(),
161+
class_id=detectron2_results["instances"]
162+
.pred_classes.cpu()
163+
.numpy()
164+
.astype(int),
165+
)
166+
167+
@classmethod
168+
def from_coco_annotations(cls, coco_annotation: dict):
169+
xyxy, class_id = [], []
170+
171+
for annotation in coco_annotation:
172+
x_min, y_min, width, height = annotation["bbox"]
173+
xyxy.append([x_min, y_min, x_min + width, y_min + height])
174+
class_id.append(annotation["category_id"])
175+
176+
return cls(xyxy=np.array(xyxy), class_id=np.array(class_id))
177+
135178
def filter(self, mask: np.ndarray, inplace: bool = False) -> Optional[Detections]:
136179
"""
137180
Filter the detections by applying a mask.
@@ -186,7 +229,9 @@ def get_anchor_coordinates(self, anchor: Position) -> np.ndarray:
186229
raise ValueError(f"{anchor} is not supported.")
187230

188231
def __getitem__(self, index: np.ndarray) -> Detections:
189-
if isinstance(index, np.ndarray) and index.dtype == bool:
232+
if isinstance(index, np.ndarray) and (
233+
index.dtype == bool or index.dtype == int
234+
):
190235
return Detections(
191236
xyxy=self.xyxy[index],
192237
confidence=self.confidence[index],
@@ -199,6 +244,17 @@ def __getitem__(self, index: np.ndarray) -> Detections:
199244
f"Detections.__getitem__ not supported for index of type {type(index)}."
200245
)
201246

247+
@property
248+
def area(self) -> np.ndarray:
249+
return (self.xyxy[:, 3] - self.xyxy[:, 1]) * (self.xyxy[:, 2] - self.xyxy[:, 0])
250+
251+
def with_nms(self, threshold: float = 0.5) -> Detections:
252+
assert (
253+
self.confidence is not None
254+
), f"Detections confidence must be given for NMS to be executed."
255+
indices = non_max_suppression(self.xyxy, self.confidence, threshold=threshold)
256+
return self[indices]
257+
202258

203259
class BoxAnnotator:
204260
def __init__(
@@ -266,7 +322,7 @@ def annotate(
266322
continue
267323

268324
text = (
269-
f"{confidence:0.2f}"
325+
f"{class_id}"
270326
if (labels is None or len(detections) != len(labels))
271327
else labels[i]
272328
)

supervision/detection/utils.py

Lines changed: 37 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -18,3 +18,40 @@ def generate_2d_mask(polygon: np.ndarray, resolution_wh: Tuple[int, int]) -> np.
1818
mask = np.zeros((height, width), dtype=np.uint8)
1919
cv2.fillPoly(mask, [polygon], color=1)
2020
return mask
21+
22+
23+
def non_max_suppression(boxes: np.ndarray, scores: np.ndarray, threshold: float):
24+
assert boxes.shape[0] == scores.shape[0]
25+
ys1 = boxes[:, 0]
26+
xs1 = boxes[:, 1]
27+
ys2 = boxes[:, 2]
28+
xs2 = boxes[:, 3]
29+
30+
areas = (ys2 - ys1) * (xs2 - xs1)
31+
scores_indexes = scores.argsort().tolist()
32+
boxes_keep_index = []
33+
while len(scores_indexes):
34+
index = scores_indexes.pop()
35+
boxes_keep_index.append(index)
36+
if not len(scores_indexes):
37+
break
38+
iou = compute_iou(
39+
boxes[index], boxes[scores_indexes], areas[index], areas[scores_indexes]
40+
)
41+
filtered_indexes = set((iou > threshold).nonzero()[0])
42+
scores_indexes = [
43+
v for (i, v) in enumerate(scores_indexes) if i not in filtered_indexes
44+
]
45+
return np.array(boxes_keep_index)
46+
47+
48+
def compute_iou(box, boxes, box_area, boxes_area):
49+
assert boxes.shape[0] == boxes_area.shape[0]
50+
ys1 = np.maximum(box[0], boxes[:, 0])
51+
xs1 = np.maximum(box[1], boxes[:, 1])
52+
ys2 = np.minimum(box[2], boxes[:, 2])
53+
xs2 = np.minimum(box[3], boxes[:, 3])
54+
intersections = np.maximum(ys2 - ys1, 0) * np.maximum(xs2 - xs1, 0)
55+
unions = box_area + boxes_area - intersections
56+
iou = intersections / unions
57+
return iou

0 commit comments

Comments
 (0)