File size: 13,933 Bytes
41eedc7 28210ca 41eedc7 28210ca 41eedc7 28210ca 90d8dcc 28210ca 90d8dcc 41eedc7 90d8dcc 41eedc7 90d8dcc 41eedc7 90d8dcc 41eedc7 90d8dcc 9229863 90d8dcc 41eedc7 28210ca 41eedc7 28210ca 41eedc7 28210ca 41eedc7 90d8dcc 41eedc7 90d8dcc |
1 2 3 4 5 6 7 8 9 10 11 12 13 14 15 16 17 18 19 20 21 22 23 24 25 26 27 28 29 30 31 32 33 34 35 36 37 38 39 40 41 42 43 44 45 46 47 48 49 50 51 52 53 54 55 56 57 58 59 60 61 62 63 64 65 66 67 68 69 70 71 72 73 74 75 76 77 78 79 80 81 82 83 84 85 86 87 88 89 90 91 92 93 94 95 96 97 98 99 100 101 102 103 104 105 106 107 108 109 110 111 112 113 114 115 116 117 118 119 120 121 122 123 124 125 126 127 128 129 130 131 132 133 134 135 136 137 138 139 140 141 142 143 144 145 146 147 148 149 150 151 152 153 154 155 156 157 158 159 160 161 162 163 164 165 166 167 168 169 170 171 172 173 174 175 176 177 178 179 180 181 182 183 184 185 186 187 188 189 190 191 192 193 194 195 196 197 198 199 200 201 202 203 204 205 206 207 208 209 210 211 212 213 214 215 216 217 218 219 220 221 222 223 224 225 226 227 228 229 230 231 232 233 234 235 236 237 238 239 240 241 242 243 244 245 246 247 248 249 250 251 252 253 254 255 256 257 258 259 260 261 262 263 264 265 266 267 268 269 270 271 272 273 274 275 276 277 278 279 280 281 282 283 284 285 286 287 288 289 290 291 292 293 294 295 296 297 298 299 300 301 302 303 304 305 306 307 308 309 310 311 312 313 314 315 316 317 318 319 320 321 322 323 324 325 326 327 328 329 330 331 332 333 334 335 336 337 338 339 340 341 342 343 344 345 346 347 348 349 350 351 352 353 354 355 356 357 358 359 360 361 362 363 364 365 366 367 368 369 370 371 372 373 374 375 376 377 378 379 380 381 382 383 384 385 386 387 388 389 390 391 392 393 394 395 396 397 398 399 400 401 402 403 404 405 406 407 408 409 410 411 412 413 414 415 416 417 418 419 420 421 422 423 424 425 426 427 428 429 430 431 432 433 434 435 436 437 438 439 440 441 442 443 444 445 446 447 448 449 450 451 452 453 454 455 456 457 458 459 460 461 462 463 464 465 466 467 468 469 470 471 472 473 |
from functools import cache
import io
import itertools
import torch
import torchvision.transforms as T
import os
import numpy as np
import seaborn as sns
from torch import nn
from torchvision.models import resnet50
from panopticapi.utils import id2rgb, rgb2id
from supervision import Detections, BoxAnnotator, MaskAnnotator
import onnx
import onnxruntime
from PIL import Image
from pathlib import Path
torch.set_grad_enabled(False)
# https://colab.research.google.com/github/facebookresearch/detr/blob/colab/notebooks/detr_demo.ipynb#scrollTo=cfCcEYjg7y46
DETR_DEMO_WEIGHTS_URI = "https://dl.fbaipublicfiles.com/detr/detr_demo-da2a99e9.pth"
TORCH_HOME = os.path.abspath(os.curdir) + "/data/cache"
ONNX_DIR = os.path.abspath(os.curdir) + "/data/onnx"
os.environ["TORCH_HOME"] = TORCH_HOME
Path(TORCH_HOME).mkdir(exist_ok=True)
Path(ONNX_DIR).mkdir(exist_ok=True)
print("Torch home:", TORCH_HOME)
# standard PyTorch mean-std input image normalization
def normalize_img(image):
transform = T.Compose(
[
T.Resize(800),
T.ToTensor(),
T.Normalize([0.485, 0.456, 0.406], [0.229, 0.224, 0.225]),
]
)
return transform(image).unsqueeze(0)
def normalize_img_800_800(image):
transform = T.Compose(
[
T.Resize((800, 800)),
T.ToTensor(),
T.Normalize([0.485, 0.456, 0.406], [0.229, 0.224, 0.225]),
]
)
return transform(image).unsqueeze(0)
# for output bounding box post-processing
def box_cxcywh_to_xyxy(x):
x_c, y_c, w, h = x.unbind(1)
b = [(x_c - 0.5 * w), (y_c - 0.5 * h), (x_c + 0.5 * w), (y_c + 0.5 * h)]
return torch.stack(b, dim=1)
def rescale_bboxes(out_bbox, size):
img_w, img_h = size
b = box_cxcywh_to_xyxy(out_bbox)
b = b * torch.tensor([img_w, img_h, img_w, img_h], dtype=torch.float32)
return b
class DETRdemo(nn.Module):
"""
Demo DETR implementation.
Demo implementation of DETR in minimal number of lines, with the
following differences wrt DETR in the paper:
* learned positional encoding (instead of sine)
* positional encoding is passed at input (instead of attention)
* fc bbox predictor (instead of MLP)
The model achieves ~40 AP on COCO val5k and runs at ~28 FPS on Tesla V100.
Only batch size 1 supported.
"""
def __init__(
self,
num_classes,
hidden_dim=256,
nheads=8,
num_encoder_layers=6,
num_decoder_layers=6,
):
super().__init__()
# create ResNet-50 backbone
self.backbone = resnet50()
del self.backbone.fc
# create conversion layer
self.conv = nn.Conv2d(2048, hidden_dim, 1)
# create a default PyTorch transformer
self.transformer = nn.Transformer(
hidden_dim, nheads, num_encoder_layers, num_decoder_layers
)
# prediction heads, one extra class for predicting non-empty slots
# note that in baseline DETR linear_bbox layer is 3-layer MLP
self.linear_class = nn.Linear(hidden_dim, num_classes + 1)
self.linear_bbox = nn.Linear(hidden_dim, 4)
# output positional encodings (object queries)
self.query_pos = nn.Parameter(torch.rand(100, hidden_dim))
# spatial positional encodings
# note that in baseline DETR we use sine positional encodings
self.row_embed = nn.Parameter(torch.rand(50, hidden_dim // 2))
self.col_embed = nn.Parameter(torch.rand(50, hidden_dim // 2))
def forward(self, inputs):
# propagate inputs through ResNet-50 up to avg-pool layer
x = self.backbone.conv1(inputs)
x = self.backbone.bn1(x)
x = self.backbone.relu(x)
x = self.backbone.maxpool(x)
x = self.backbone.layer1(x)
x = self.backbone.layer2(x)
x = self.backbone.layer3(x)
x = self.backbone.layer4(x)
# convert from 2048 to 256 feature planes for the transformer
h = self.conv(x)
# construct positional encodings
H, W = h.shape[-2:]
pos = (
torch.cat(
[
self.col_embed[:W].unsqueeze(0).repeat(H, 1, 1),
self.row_embed[:H].unsqueeze(1).repeat(1, W, 1),
],
dim=-1,
)
.flatten(0, 1)
.unsqueeze(1)
)
# propagate through the transformer
h = self.transformer(
pos + 0.1 * h.flatten(2).permute(2, 0, 1), self.query_pos.unsqueeze(1)
).transpose(0, 1)
# finally project transformer outputs to class labels and bounding boxes
return {
"pred_logits": self.linear_class(h),
"pred_boxes": self.linear_bbox(h).sigmoid(),
}
class SimpleDetr:
@cache
def __init__(self):
self.model = DETRdemo(num_classes=91)
state_dict = torch.hub.load_state_dict_from_url(
url=DETR_DEMO_WEIGHTS_URI,
map_location="cpu",
check_hash=True,
)
self.model.load_state_dict(state_dict)
self.model.eval()
self.box_annotator: BoxAnnotator = BoxAnnotator()
def detect(self, image, conf):
# mean-std normalize the input image (batch-size: 1)
img = normalize_img(image)
# demo model only support by default images with aspect ratio between 0.5 and 2
# if you want to use images with an aspect ratio outside this range
# rescale your image so that the maximum size is at most 1333 for best results
assert (
img.shape[-2] <= 1600 and img.shape[-1] <= 1600
), "demo model only supports images up to 1600 pixels on each side"
# propagate through the model
outputs = self.model(img)
# keep only predictions with 0.7+ confidence
scores = outputs["pred_logits"].softmax(-1)[0, :, :-1]
keep = scores.max(-1).values > conf
# convert boxes from [0; 1] to image scales
bboxes_scaled = rescale_bboxes(outputs["pred_boxes"][0, keep], image.size)
probas = scores[keep]
class_id = []
confidence = []
for prob in probas:
cls_id = prob.argmax()
c = prob[cls_id]
class_id.append(int(cls_id))
confidence.append(float(c))
print(class_id, confidence)
detections = Detections(
xyxy=bboxes_scaled.cpu().detach().numpy(),
class_id=np.array(class_id),
confidence=np.array(confidence),
)
annotated = self.box_annotator.annotate(
scene=np.array(image),
skip_label=False,
detections=detections,
labels=[
f"{CLASSES[cls_id]} {conf:.2f}"
for cls_id, conf in zip(detections.class_id, detections.confidence)
],
)
return annotated
def export(self):
model_path = f"{ONNX_DIR}/detr_simple_demo_onnx.onnx"
dummy_image = torch.ones(1, 3, 800, 800, device="cpu")
input_names = ["inputs"]
output_names = ["pred_logits", "pred_boxes"]
torch.onnx.export(
self.model,
dummy_image,
model_path,
input_names=input_names,
output_names=output_names,
# dynamic_axes={input_names[0]: {0: "batch_size", 2: "height", 3: "width"}}, #!TODO
export_params=True,
training=torch.onnx.TrainingMode.EVAL,
opset_version=14,
)
onnx_model = onnx.load(model_path)
# Check the model
try:
onnx.checker.check_model(onnx_model)
except onnx.checker.ValidationError as e:
print(f"The model is invalid: {e}")
else:
print("The model is valid!")
return model_path
class SimpleDetrOnnx:
@cache
def __init__(self):
self.box_annotator: BoxAnnotator = BoxAnnotator()
onnx_sess_opts = onnxruntime.SessionOptions()
onnx_sess_opts.graph_optimization_level = (
onnxruntime.GraphOptimizationLevel.ORT_ENABLE_ALL
# onnxruntime.GraphOptimizationLevel.ORT_DISABLE_ALL
)
onnx_sess_opts.enable_mem_pattern = True
onnx_sess_opts.enable_cpu_mem_arena = True
self.ort_session = onnxruntime.InferenceSession(
f"{ONNX_DIR}/detr_simple_demo_onnx.onnx",
sess_options=onnx_sess_opts,
providers=[
"CUDAExecutionProvider",
"CoreMLExecutionProvider",
"CPUExecutionProvider",
],
)
self.classes = {}
self.metadata = self.ort_session.get_modelmeta()
self.providers = self.ort_session.get_providers()
print(f"[OnnxRuntime] Providers:{self.providers}")
print(
f"[OnnxRuntime] medatadata: {self.metadata.custom_metadata_map} {type(self.metadata.custom_metadata_map)}"
)
def detect(self, image, conf):
# dummy_image = np.ones((1, 3, 600, 800), dtype=np.float32)
im = normalize_img_800_800(image).numpy()
print("SHAPE", im.shape)
ort_inputs = {self.ort_session.get_inputs()[0].name: im}
outputs = self.ort_session.run(None, ort_inputs)
pred_logits = torch.tensor(
outputs[0]
) # conversion to torch for simplicity (softmax etc)
pred_boxes = torch.tensor(outputs[1])
scores = pred_logits.softmax(-1)[0, :, :-1]
keep = scores.max(-1).values > conf
bboxes_scaled = rescale_bboxes(pred_boxes[0, keep], image.size)
probas = scores[keep]
class_id = []
confidence = []
for prob in probas:
cls_id = prob.argmax()
c = prob[cls_id]
class_id.append(int(cls_id))
confidence.append(float(c))
print(class_id, confidence)
detections = Detections(
xyxy=bboxes_scaled.cpu().detach().numpy(),
class_id=np.array(class_id),
confidence=np.array(confidence),
)
annotated = self.box_annotator.annotate(
scene=np.array(image),
skip_label=False,
detections=detections,
labels=[
f"{CLASSES[cls_id]} {conf:.2f}"
for cls_id, conf in zip(detections.class_id, detections.confidence)
],
)
return annotated
class PanopticDetrResenet101:
@cache
def __init__(self):
self.model, self.postprocessor = torch.hub.load(
"facebookresearch/detr",
"detr_resnet101_panoptic",
pretrained=True,
return_postprocessor=True,
num_classes=250,
)
self.model.eval()
def detect(self, image, conf):
# mean-std normalize the input image (batch-size: 1)
img = normalize_img(image)
outputs = self.model(img)
result = self.postprocessor(
outputs, torch.as_tensor(img.shape[-2:]).unsqueeze(0)
)[0]
print(result.keys())
palette = itertools.cycle(sns.color_palette())
# The segmentation is stored in a special-format png
panoptic_seg = Image.open(io.BytesIO(result["png_string"]))
panoptic_seg = np.array(panoptic_seg, dtype=np.uint8).copy()
# We retrieve the ids corresponding to each mask
panoptic_seg_id = rgb2id(panoptic_seg)
# Finally we color each mask individually
panoptic_seg[:, :, :] = 0
for id in range(panoptic_seg_id.max() + 1):
panoptic_seg[panoptic_seg_id == id] = np.asarray(next(palette)) * 255
return panoptic_seg
def export(self):
model_path = f"{ONNX_DIR}/detr_resnet101_panoptic.onnx"
dummy_image = torch.ones(1, 3, 800, 800, device="cpu")
input_names = ["inputs"]
output_names = ["pred_logits", "pred_boxes", "pred_masks"]
torch.onnx.export(
self.model,
dummy_image,
model_path,
input_names=input_names,
output_names=output_names,
export_params=True,
training=torch.onnx.TrainingMode.EVAL,
opset_version=14,
)
onnx_model = onnx.load(model_path)
# Check the model
try:
onnx.checker.check_model(onnx_model)
except onnx.checker.ValidationError as e:
print(f"The model is invalid: {e}")
else:
print("The model is valid!")
return model_path
# COCO classes
CLASSES = [
"N/A",
"person",
"bicycle",
"car",
"motorcycle",
"airplane",
"bus",
"train",
"truck",
"boat",
"traffic light",
"fire hydrant",
"N/A",
"stop sign",
"parking meter",
"bench",
"bird",
"cat",
"dog",
"horse",
"sheep",
"cow",
"elephant",
"bear",
"zebra",
"giraffe",
"N/A",
"backpack",
"umbrella",
"N/A",
"N/A",
"handbag",
"tie",
"suitcase",
"frisbee",
"skis",
"snowboard",
"sports ball",
"kite",
"baseball bat",
"baseball glove",
"skateboard",
"surfboard",
"tennis racket",
"bottle",
"N/A",
"wine glass",
"cup",
"fork",
"knife",
"spoon",
"bowl",
"banana",
"apple",
"sandwich",
"orange",
"broccoli",
"carrot",
"hot dog",
"pizza",
"donut",
"cake",
"chair",
"couch",
"potted plant",
"bed",
"N/A",
"dining table",
"N/A",
"N/A",
"toilet",
"N/A",
"tv",
"laptop",
"mouse",
"remote",
"keyboard",
"cell phone",
"microwave",
"oven",
"toaster",
"sink",
"refrigerator",
"N/A",
"book",
"clock",
"vase",
"scissors",
"teddy bear",
"hair drier",
"toothbrush",
]
# model = SimpleDetr()
# model.export()
|