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

Add TensorFlow and TFLite export #1127

Merged
merged 71 commits into from
Aug 17, 2021
Merged
Show file tree
Hide file tree
Changes from 61 commits
Commits
Show all changes
71 commits
Select commit Hold shift + click to select a range
5b6528c
Add models/tf.py for TensorFlow and TFLite export
zldrobit Oct 9, 2020
a30dad4
Set auto=False for int8 calibration
zldrobit Oct 9, 2020
f0cb6e2
Update requirements.txt for TensorFlow and TFLite export
zldrobit Oct 20, 2020
ce73d3d
Read anchors directly from PyTorch weights
zldrobit Oct 23, 2020
d101f7e
Add --tf-nms to append NMS in TensorFlow SavedModel and GraphDef export
zldrobit Nov 2, 2020
55fb2eb
Remove check_anchor_order, check_file, set_logging from import
zldrobit Nov 16, 2020
3dd69d9
Reformat code and optimize imports
glenn-jocher Nov 28, 2020
e8a9ad2
Autodownload model and check cfg
glenn-jocher Nov 28, 2020
efbb853
update --source path, img-size to 320, single output
glenn-jocher Nov 28, 2020
aed53ce
Adjust representative_dataset
glenn-jocher Nov 28, 2020
ccb2336
Put representative dataset in tfl_int8 block
zldrobit Nov 30, 2020
9f893c8
detect.py TF inference
glenn-jocher Dec 2, 2020
d9fad06
Merge remote-tracking branch 'origin/tf-only-export' into tf-only-export
glenn-jocher Dec 2, 2020
49a9e05
weights to string
glenn-jocher Dec 2, 2020
1867bb4
weights to string
glenn-jocher Dec 2, 2020
4eed608
cleanup tf.py
glenn-jocher Dec 4, 2020
8ba2ca9
Add --dynamic-batch-size
zldrobit Dec 22, 2020
4d9104a
Add xywh normalization to reduce calibration error
zldrobit Dec 22, 2020
ae9bce8
Merge branch 'master' into tf-only-export
glenn-jocher Dec 22, 2020
cabb802
Update requirements.txt
zldrobit Dec 23, 2020
a5967f8
Fix imports
zldrobit Dec 24, 2020
dbc7f71
Add models/tf.py for TensorFlow and TFLite export
zldrobit Oct 9, 2020
565e620
Set auto=False for int8 calibration
zldrobit Oct 9, 2020
fc26561
Update requirements.txt for TensorFlow and TFLite export
zldrobit Oct 20, 2020
05cc389
Read anchors directly from PyTorch weights
zldrobit Oct 23, 2020
817fcf8
Add --tf-nms to append NMS in TensorFlow SavedModel and GraphDef export
zldrobit Nov 2, 2020
9a4ce5e
Remove check_anchor_order, check_file, set_logging from import
zldrobit Nov 16, 2020
719479e
Reformat code and optimize imports
glenn-jocher Nov 28, 2020
243fa7f
Autodownload model and check cfg
glenn-jocher Nov 28, 2020
3b6cf12
update --source path, img-size to 320, single output
glenn-jocher Nov 28, 2020
5a5f949
Adjust representative_dataset
glenn-jocher Nov 28, 2020
00d50fd
detect.py TF inference
glenn-jocher Dec 2, 2020
061907b
Put representative dataset in tfl_int8 block
zldrobit Nov 30, 2020
ca4550b
weights to string
glenn-jocher Dec 2, 2020
5e04c5c
weights to string
glenn-jocher Dec 2, 2020
9121a87
cleanup tf.py
glenn-jocher Dec 4, 2020
e9bc606
Add --dynamic-batch-size
zldrobit Dec 22, 2020
dacd8af
Add xywh normalization to reduce calibration error
zldrobit Dec 22, 2020
b492af9
Update requirements.txt
zldrobit Dec 23, 2020
fbf5a45
Fix imports
zldrobit Dec 24, 2020
c4cfbd9
Merge branch 'tf-only-export' of https://github.com/zldrobit/yolov5 i…
glenn-jocher Jan 8, 2021
c761637
implement C3() and SiLU()
glenn-jocher Jan 8, 2021
36ed3cd
Fix reshape dim to support dynamic batching
zldrobit Feb 2, 2021
aad9e24
Merge branch 'master' into tf-only-export
glenn-jocher Feb 4, 2021
8ec975e
Merge branch 'master' into tf-only-export
glenn-jocher Feb 6, 2021
b0fa5a3
Add epsilon argument in tf_BN, which is different between TF and PT
zldrobit Mar 11, 2021
710bf56
Set stride to None if not using PyTorch, and do not warmup without Py…
zldrobit Mar 12, 2021
c45ceef
Add list support in check_img_size()
zldrobit Mar 23, 2021
0d39b24
Add list input support in detect.py
zldrobit Mar 23, 2021
47da942
merge ultralytics:master
glenn-jocher Mar 25, 2021
8cb7032
merge ultralytics:master
glenn-jocher Mar 28, 2021
4e1485b
sys.path.append('./') to run from yolov5/
glenn-jocher Mar 28, 2021
e4e6d6f
Add int8 quantization support for TensorFlow 2.5
zldrobit Apr 3, 2021
aafe224
Add get_coco128.sh
zldrobit May 6, 2021
d3be281
Remove --no-tfl-detect in models/tf.py (Use tf-android-tfl-detect bra…
zldrobit May 6, 2021
9ca5d7a
Merge branch 'develop' into tf-only-export
glenn-jocher Jun 2, 2021
215c865
Update requirements.txt
glenn-jocher Jun 2, 2021
a2867da
Replace torch.load() with attempt_load()
zldrobit Jun 7, 2021
86768d1
Update requirements.txt
zldrobit Jun 7, 2021
972c6a2
Add --tf-raw-resize to set half_pixel_centers=False
zldrobit Jun 11, 2021
eed4980
Add --agnostic-nms for TF class-agnostic NMS
zldrobit Jun 18, 2021
34b7d67
Merge master
glenn-jocher Aug 16, 2021
10eebf2
Cleanup after merge
glenn-jocher Aug 16, 2021
e3aa755
Cleanup2 after merge
glenn-jocher Aug 16, 2021
27fc39b
Cleanup3 after merge
glenn-jocher Aug 16, 2021
48a6bf8
Add tf.py docstring with credit and usage
glenn-jocher Aug 16, 2021
4e12c70
pb saved_model and tflite use only one model in detect.py
zldrobit Aug 17, 2021
ea0274f
Add use cases in docstring of tf.py
zldrobit Aug 17, 2021
2f63c6d
Remove redundant `stride` definition
glenn-jocher Aug 17, 2021
c3f46c3
Remove keras direct import
glenn-jocher Aug 17, 2021
3133381
Fix `check_requirements(('tensorflow>=2.4.1',))`
glenn-jocher Aug 17, 2021
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
101 changes: 88 additions & 13 deletions detect.py
Original file line number Diff line number Diff line change
Expand Up @@ -5,6 +5,9 @@
import cv2
import torch
import torch.backends.cudnn as cudnn
import yaml
from numpy import random
import numpy as np

from models.experimental import attempt_load
from utils.datasets import LoadStreams, LoadImages
Expand All @@ -31,12 +34,57 @@ def detect(opt):
half = device.type != 'cpu' # half precision only supported on CUDA

# Load model
model = attempt_load(weights, map_location=device) # load FP32 model
stride = int(model.stride.max()) # model stride
imgsz = check_img_size(imgsz, s=stride) # check img_size
names = model.module.names if hasattr(model, 'module') else model.names # get class names
if half:
model.half() # to FP16
weights = weights[0] if isinstance(weights, list) else weights
suffix = Path(weights).suffix
if suffix == '.pt':
backend = 'pytorch'
model = attempt_load(weights, map_location=device) # load FP32 model
stride = int(model.stride.max()) # model stride
imgsz = check_img_size(imgsz, s=stride) # check img_size
names = model.module.names if hasattr(model, 'module') else model.names # class names
if half:
model.half() # to FP16
else:
import tensorflow as tf
from tensorflow import keras

stride = None
with open('data/coco.yaml') as f:
names = yaml.load(f, Loader=yaml.FullLoader)['names'] # class names (assume COCO)

if suffix == '.pb':
backend = 'graph_def'

# https://www.tensorflow.org/guide/migrate#a_graphpb_or_graphpbtxt
# https://github.com/leimao/Frozen_Graph_TensorFlow
def wrap_frozen_graph(graph_def, inputs, outputs):
def _imports_graph_def():
tf.compat.v1.import_graph_def(graph_def, name="")

wrapped_import = tf.compat.v1.wrap_function(_imports_graph_def, [])
import_graph = wrapped_import.graph
return wrapped_import.prune(
tf.nest.map_structure(import_graph.as_graph_element, inputs),
tf.nest.map_structure(import_graph.as_graph_element, outputs))

graph = tf.Graph()
graph_def = graph.as_graph_def()
graph_def.ParseFromString(open(weights, 'rb').read())
frozen_func = wrap_frozen_graph(graph_def=graph_def, inputs="x:0", outputs="Identity:0")

elif suffix == '.tflite':
backend = 'tflite'
# Load TFLite model and allocate tensors
interpreter = tf.lite.Interpreter(model_path=weights)
interpreter.allocate_tensors()

# Get input and output tensors
input_details = interpreter.get_input_details()
output_details = interpreter.get_output_details()

else:
backend = 'saved_model'
model = keras.models.load_model(weights)

# Second-stage classifier
classify = False
Expand All @@ -49,24 +97,49 @@ def detect(opt):
if webcam:
view_img = check_imshow()
cudnn.benchmark = True # set True to speed up constant image size inference
dataset = LoadStreams(source, img_size=imgsz, stride=stride)
dataset = LoadStreams(source, img_size=imgsz, stride=stride, auto=backend == 'pytorch')
else:
dataset = LoadImages(source, img_size=imgsz, stride=stride)
dataset = LoadImages(source, img_size=imgsz, stride=stride, auto=backend == 'pytorch')

# Run inference
if device.type != 'cpu':
model(torch.zeros(1, 3, imgsz, imgsz).to(device).type_as(next(model.parameters()))) # run once
if device.type != 'cpu' and backend == 'pytorch':
model(torch.zeros(1, 3, *imgsz).to(device).type_as(next(model.parameters()))) # run once
t0 = time.time()
for path, img, im0s, vid_cap in dataset:
img = torch.from_numpy(img).to(device)
img = img.half() if half else img.float() # uint8 to fp16/32
img = img.half() if half and backend == 'pytorch' else img.float() # uint8 to fp16/32
img /= 255.0 # 0 - 255 to 0.0 - 1.0
if img.ndimension() == 3:
img = img.unsqueeze(0)

# Inference
t1 = time_synchronized()
pred = model(img, augment=opt.augment)[0]
if backend == 'pytorch':
pred = model(img, augment=opt.augment)[0]
else:
if backend == 'saved_model':
pred = model(img.permute(0, 2, 3, 1).cpu().numpy(), training=False).numpy()
elif backend == 'graph_def':
pred = frozen_func(x=tf.constant(img.permute(0, 2, 3, 1).cpu().numpy())).numpy()
elif backend == 'tflite':
input_data = img.permute(0, 2, 3, 1).cpu().numpy()
if opt.tfl_int8:
scale, zero_point = input_details[0]['quantization']
input_data = input_data / scale + zero_point
input_data = input_data.astype(np.uint8)
interpreter.set_tensor(input_details[0]['index'], input_data)
interpreter.invoke()
pred = interpreter.get_tensor(output_details[0]['index'])
if opt.tfl_int8:
scale, zero_point = output_details[0]['quantization']
pred = pred.astype(np.float32)
pred = (pred - zero_point) * scale
# Denormalize xywh
pred[..., 0] *= imgsz[1] # x
pred[..., 1] *= imgsz[0] # y
pred[..., 2] *= imgsz[1] # w
pred[..., 3] *= imgsz[0] # h
pred = torch.tensor(pred)

# Apply NMS
pred = non_max_suppression(pred, opt.conf_thres, opt.iou_thres, opt.classes, opt.agnostic_nms,
Expand Down Expand Up @@ -152,7 +225,7 @@ def detect(opt):
parser = argparse.ArgumentParser()
parser.add_argument('--weights', nargs='+', type=str, default='yolov5s.pt', help='model.pt path(s)')
parser.add_argument('--source', type=str, default='data/images', help='source') # file/folder, 0 for webcam
parser.add_argument('--img-size', type=int, default=640, help='inference size (pixels)')
parser.add_argument('--img-size', nargs='+', type=int, default=[320, 320], help='image size') # height, width
parser.add_argument('--conf-thres', type=float, default=0.25, help='object confidence threshold')
parser.add_argument('--iou-thres', type=float, default=0.45, help='IOU threshold for NMS')
parser.add_argument('--max-det', type=int, default=1000, help='maximum number of detections per image')
Expand All @@ -169,10 +242,12 @@ def detect(opt):
parser.add_argument('--project', default='runs/detect', help='save results to project/name')
parser.add_argument('--name', default='exp', help='save results to project/name')
parser.add_argument('--exist-ok', action='store_true', help='existing project/name ok, do not increment')
parser.add_argument('--tfl-int8', action='store_true', help='use int8 quantized TFLite model')
parser.add_argument('--line-thickness', default=3, type=int, help='bounding box thickness (pixels)')
parser.add_argument('--hide-labels', default=False, action='store_true', help='hide labels')
parser.add_argument('--hide-conf', default=False, action='store_true', help='hide confidences')
opt = parser.parse_args()
opt.img_size *= 2 if len(opt.img_size) == 1 else 1 # expand
print(opt)
check_requirements(exclude=('tensorboard', 'pycocotools', 'thop'))

Expand Down
8 changes: 6 additions & 2 deletions models/experimental.py
Original file line number Diff line number Diff line change
Expand Up @@ -110,14 +110,18 @@ def forward(self, x, augment=False):
return y, None # inference, train output


def attempt_load(weights, map_location=None, inplace=True):
def attempt_load(weights, map_location=None, inplace=True, fuse=True):
from models.yolo import Detect, Model

# Loads an ensemble of models weights=[a,b,c] or a single model weights=[a] or weights=a
model = Ensemble()
for w in weights if isinstance(weights, list) else [weights]:
ckpt = torch.load(attempt_download(w), map_location=map_location) # load
model.append(ckpt['ema' if ckpt.get('ema') else 'model'].float().fuse().eval()) # FP32 model
if fuse:
model.append(ckpt['ema' if ckpt.get('ema') else 'model'].float().fuse().eval()) # FP32 model
else:
model.append(ckpt['ema' if ckpt.get('ema') else 'model'].float().eval()) # without layer fuse


# Compatibility updates
for m in model.modules():
Expand Down
Loading