Skip to content

Commit 33f48a9

Browse files
authored
feat(data): support cache ram of COCO dataset (Megvii-BaseDetection#1562)
feat(data): support cache ram of COCO dataset
1 parent 11c2a1f commit 33f48a9

File tree

7 files changed

+152
-79
lines changed

7 files changed

+152
-79
lines changed

requirements.txt

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -8,6 +8,7 @@ torchvision
88
thop
99
ninja
1010
tabulate
11+
psutil
1112

1213
# verified versions
1314
# pycocotools corresponds to https://github.com/ppwwyyxx/cocoapi

setup.cfg

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -3,7 +3,7 @@ line_length = 100
33
multi_line_output = 3
44
balanced_wrapping = True
55
known_standard_library = setuptools
6-
known_third_party = tqdm,loguru,tabulate
6+
known_third_party = tqdm,loguru,tabulate,psutil
77
known_data_processing = cv2,numpy,scipy,PIL,matplotlib
88
known_datasets = pycocotools
99
known_deeplearning = torch,torchvision,caffe2,onnx,apex,timm,thop,torch2trt,tensorrt,openvino,onnxruntime

tools/train.py

Lines changed: 7 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -67,10 +67,10 @@ def make_parser():
6767
)
6868
parser.add_argument(
6969
"--cache",
70-
dest="cache",
71-
default=False,
72-
action="store_true",
73-
help="Caching imgs to RAM for fast training.",
70+
type=str,
71+
nargs="?",
72+
const="ram",
73+
help="Caching imgs to ram/disk for fast training.",
7474
)
7575
parser.add_argument(
7676
"-o",
@@ -130,6 +130,9 @@ def main(exp: Exp, args):
130130
num_gpu = get_num_devices() if args.devices is None else args.devices
131131
assert num_gpu <= get_num_devices()
132132

133+
if args.cache is not None:
134+
exp.create_cache_dataset(args.cache)
135+
133136
dist_url = "auto" if args.dist_url is None else args.dist_url
134137
launch(
135138
main,

yolox/core/trainer.py

Lines changed: 5 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -26,6 +26,7 @@
2626
gpu_mem_usage,
2727
is_parallel,
2828
load_ckpt,
29+
mem_usage,
2930
occupy_mem,
3031
save_checkpoint,
3132
setup_logger,
@@ -250,10 +251,12 @@ def after_iter(self):
250251
["{}: {:.3f}s".format(k, v.avg) for k, v in time_meter.items()]
251252
)
252253

254+
mem_str = "gpu mem: {:.0f}Mb, mem: {:.1f}Gb".format(gpu_mem_usage(), mem_usage())
255+
253256
logger.info(
254-
"{}, mem: {:.0f}Mb, {}, {}, lr: {:.3e}".format(
257+
"{}, {}, {}, {}, lr: {:.3e}".format(
255258
progress_str,
256-
gpu_mem_usage(),
259+
mem_str,
257260
time_str,
258261
loss_str,
259262
self.meter["lr"].latest,

yolox/data/datasets/coco.py

Lines changed: 84 additions & 58 deletions
Original file line numberDiff line numberDiff line change
@@ -1,9 +1,13 @@
11
#!/usr/bin/env python3
22
# -*- coding:utf-8 -*-
33
# Copyright (c) Megvii, Inc. and its affiliates.
4-
4+
import copy
55
import os
6+
import random
7+
from multiprocessing.pool import ThreadPool
8+
import psutil
69
from loguru import logger
10+
from tqdm import tqdm
711

812
import cv2
913
import numpy as np
@@ -45,6 +49,7 @@ def __init__(
4549
img_size=(416, 416),
4650
preproc=None,
4751
cache=False,
52+
cache_type="ram",
4853
):
4954
"""
5055
COCO dataset initialization. Annotation data are read into memory by COCO API.
@@ -64,74 +69,95 @@ def __init__(
6469
self.coco = COCO(os.path.join(self.data_dir, "annotations", self.json_file))
6570
remove_useless_info(self.coco)
6671
self.ids = self.coco.getImgIds()
72+
self.num_imgs = len(self.ids)
6773
self.class_ids = sorted(self.coco.getCatIds())
6874
self.cats = self.coco.loadCats(self.coco.getCatIds())
6975
self._classes = tuple([c["name"] for c in self.cats])
70-
self.imgs = None
7176
self.name = name
7277
self.img_size = img_size
7378
self.preproc = preproc
7479
self.annotations = self._load_coco_annotations()
75-
if cache:
80+
self.imgs = None
81+
self.cache = cache
82+
self.cache_type = cache_type
83+
84+
if self.cache:
7685
self._cache_images()
7786

78-
def __len__(self):
79-
return len(self.ids)
87+
def _cache_images(self):
88+
mem = psutil.virtual_memory()
89+
mem_required = self.cal_cache_ram()
90+
gb = 1 << 30
8091

81-
def __del__(self):
82-
del self.imgs
92+
if self.cache_type == "ram" and mem_required > mem.available:
93+
self.cache = False
94+
else:
95+
logger.info(
96+
f"{mem_required / gb:.1f}GB RAM required, "
97+
f"{mem.available / gb:.1f}/{mem.total / gb:.1f}GB RAM available, "
98+
f"Since the first thing we do is cache, "
99+
f"there is no guarantee that the remaining memory space is sufficient"
100+
)
83101

84-
def _load_coco_annotations(self):
85-
return [self.load_anno_from_ids(_ids) for _ids in self.ids]
102+
if self.cache and self.imgs is None:
103+
if self.cache_type == 'ram':
104+
self.imgs = [None] * self.num_imgs
105+
logger.info("You are using cached images in RAM to accelerate training!")
106+
else: # 'disk'
107+
self.cache_dir = os.path.join(
108+
self.data_dir,
109+
f"{self.name}_cache{self.img_size[0]}x{self.img_size[1]}"
110+
)
111+
if not os.path.exists(self.cache_dir):
112+
os.mkdir(self.cache_dir)
113+
logger.warning(
114+
f"\n*******************************************************************\n"
115+
f"You are using cached images in DISK to accelerate training.\n"
116+
f"This requires large DISK space.\n"
117+
f"Make sure you have {mem_required / gb:.1f} "
118+
f"available DISK space for training COCO.\n"
119+
f"*******************************************************************\\n"
120+
)
121+
else:
122+
logger.info("Found disk cache!")
123+
return
86124

87-
def _cache_images(self):
88-
logger.warning(
89-
"\n********************************************************************************\n"
90-
"You are using cached images in RAM to accelerate training.\n"
91-
"This requires large system RAM.\n"
92-
"Make sure you have 200G+ RAM and 136G available disk space for training COCO.\n"
93-
"********************************************************************************\n"
94-
)
95-
max_h = self.img_size[0]
96-
max_w = self.img_size[1]
97-
cache_file = os.path.join(self.data_dir, f"img_resized_cache_{self.name}.array")
98-
if not os.path.exists(cache_file):
99125
logger.info(
100-
"Caching images for the first time. This might take about 20 minutes for COCO"
126+
"Caching images for the first time. "
127+
"This might take about 15 minutes for COCO"
101128
)
102-
self.imgs = np.memmap(
103-
cache_file,
104-
shape=(len(self.ids), max_h, max_w, 3),
105-
dtype=np.uint8,
106-
mode="w+",
107-
)
108-
from tqdm import tqdm
109-
from multiprocessing.pool import ThreadPool
110129

111-
NUM_THREADs = min(8, os.cpu_count())
112-
loaded_images = ThreadPool(NUM_THREADs).imap(
113-
lambda x: self.load_resized_img(x),
114-
range(len(self.annotations)),
115-
)
116-
pbar = tqdm(enumerate(loaded_images), total=len(self.annotations))
117-
for k, out in pbar:
118-
self.imgs[k][: out.shape[0], : out.shape[1], :] = out.copy()
119-
self.imgs.flush()
130+
num_threads = min(8, max(1, os.cpu_count() - 1))
131+
b = 0
132+
load_imgs = ThreadPool(num_threads).imap(self.load_resized_img, range(self.num_imgs))
133+
pbar = tqdm(enumerate(load_imgs), total=self.num_imgs)
134+
for i, x in pbar: # x = self.load_resized_img(self, i)
135+
if self.cache_type == 'ram':
136+
self.imgs[i] = x
137+
else: # 'disk'
138+
cache_filename = f'{self.annotations[i]["filename"].split(".")[0]}.npy'
139+
np.save(os.path.join(self.cache_dir, cache_filename), x)
140+
b += x.nbytes
141+
pbar.desc = f'Caching images ({b / gb:.1f}/{mem_required / gb:.1f}GB {self.cache})'
120142
pbar.close()
121-
else:
122-
logger.warning(
123-
"You are using cached imgs! Make sure your dataset is not changed!!\n"
124-
"Everytime the self.input_size is changed in your exp file, you need to delete\n"
125-
"the cached data and re-generate them.\n"
126-
)
127143

128-
logger.info("Loading cached imgs...")
129-
self.imgs = np.memmap(
130-
cache_file,
131-
shape=(len(self.ids), max_h, max_w, 3),
132-
dtype=np.uint8,
133-
mode="r+",
134-
)
144+
def cal_cache_ram(self):
145+
cache_bytes = 0
146+
num_samples = min(self.num_imgs, 32)
147+
for _ in range(num_samples):
148+
img = self.load_resized_img(random.randint(0, self.num_imgs - 1))
149+
cache_bytes += img.nbytes
150+
mem_required = cache_bytes * self.num_imgs / num_samples
151+
return mem_required
152+
153+
def __len__(self):
154+
return self.num_imgs
155+
156+
def __del__(self):
157+
del self.imgs
158+
159+
def _load_coco_annotations(self):
160+
return [self.load_anno_from_ids(_ids) for _ids in self.ids]
135161

136162
def load_anno_from_ids(self, id_):
137163
im_ann = self.coco.loadImgs(id_)[0]
@@ -152,7 +178,6 @@ def load_anno_from_ids(self, id_):
152178
num_objs = len(objs)
153179

154180
res = np.zeros((num_objs, 5))
155-
156181
for ix, obj in enumerate(objs):
157182
cls = self.class_ids.index(obj["category_id"])
158183
res[ix, 0:4] = obj["clean_bbox"]
@@ -197,15 +222,16 @@ def load_image(self, index):
197222

198223
def pull_item(self, index):
199224
id_ = self.ids[index]
225+
label, origin_image_size, _, filename = self.annotations[index]
200226

201-
res, img_info, resized_info, _ = self.annotations[index]
202-
if self.imgs is not None:
203-
pad_img = self.imgs[index]
204-
img = pad_img[: resized_info[0], : resized_info[1], :].copy()
227+
if self.cache_type == 'ram':
228+
img = self.imgs[index]
229+
elif self.cache_type == 'disk':
230+
img = np.load(os.path.join(self.cache_dir, f"{filename.split('.')[0]}.npy"))
205231
else:
206232
img = self.load_resized_img(index)
207233

208-
return img, res.copy(), img_info, np.array([id_])
234+
return copy.deepcopy(img), copy.deepcopy(label), origin_image_size, np.array([id_])
209235

210236
@Dataset.mosaic_getitem
211237
def __getitem__(self, index):

yolox/exp/yolox_base.py

Lines changed: 43 additions & 14 deletions
Original file line numberDiff line numberDiff line change
@@ -106,6 +106,23 @@ def __init__(self):
106106
self.test_conf = 0.01
107107
# nms threshold
108108
self.nmsthre = 0.65
109+
self.cache_dataset = None
110+
self.dataset = None
111+
112+
def create_cache_dataset(self, cache_type: str = "ram"):
113+
from yolox.data import COCODataset, TrainTransform
114+
self.cache_dataset = COCODataset(
115+
data_dir=self.data_dir,
116+
json_file=self.train_ann,
117+
img_size=self.input_size,
118+
preproc=TrainTransform(
119+
max_labels=50,
120+
flip_prob=self.flip_prob,
121+
hsv_prob=self.hsv_prob
122+
),
123+
cache=True,
124+
cache_type=cache_type,
125+
)
109126

110127
def get_model(self):
111128
from yolox.models import YOLOX, YOLOPAFPN, YOLOXHead
@@ -127,7 +144,16 @@ def init_yolo(M):
127144
self.model.train()
128145
return self.model
129146

130-
def get_data_loader(self, batch_size, is_distributed, no_aug=False, cache_img=False):
147+
def get_data_loader(self, batch_size, is_distributed, no_aug=False, cache_img: str = None):
148+
"""
149+
Get dataloader according to cache_img parameter.
150+
Args:
151+
no_aug (bool, optional): Whether to turn off mosaic data enhancement. Defaults to False.
152+
cache_img (str, optional): cache_img is equivalent to cache_type. Defaults to None.
153+
"ram" : Caching imgs to ram for fast training.
154+
"disk": Caching imgs to disk for fast training.
155+
None: Do not use cache, in this case cache_data is also None.
156+
"""
131157
from yolox.data import (
132158
COCODataset,
133159
TrainTransform,
@@ -140,18 +166,23 @@ def get_data_loader(self, batch_size, is_distributed, no_aug=False, cache_img=Fa
140166
from yolox.utils import wait_for_the_master
141167

142168
with wait_for_the_master():
143-
dataset = COCODataset(
144-
data_dir=self.data_dir,
145-
json_file=self.train_ann,
146-
img_size=self.input_size,
147-
preproc=TrainTransform(
148-
max_labels=50,
149-
flip_prob=self.flip_prob,
150-
hsv_prob=self.hsv_prob),
151-
cache=cache_img,
152-
)
169+
if self.cache_dataset is None:
170+
assert cache_img is None, "cache is True, but cache_dataset is None"
171+
dataset = COCODataset(
172+
data_dir=self.data_dir,
173+
json_file=self.train_ann,
174+
img_size=self.input_size,
175+
preproc=TrainTransform(
176+
max_labels=50,
177+
flip_prob=self.flip_prob,
178+
hsv_prob=self.hsv_prob),
179+
cache=False,
180+
cache_type=cache_img,
181+
)
182+
else:
183+
dataset = self.cache_dataset
153184

154-
dataset = MosaicDetection(
185+
self.dataset = MosaicDetection(
155186
dataset,
156187
mosaic=not no_aug,
157188
img_size=self.input_size,
@@ -169,8 +200,6 @@ def get_data_loader(self, batch_size, is_distributed, no_aug=False, cache_img=Fa
169200
mixup_prob=self.mixup_prob,
170201
)
171202

172-
self.dataset = dataset
173-
174203
if is_distributed:
175204
batch_size = batch_size // dist.get_world_size()
176205

yolox/utils/metric.py

Lines changed: 11 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -5,6 +5,7 @@
55
import os
66
import time
77
from collections import defaultdict, deque
8+
import psutil
89

910
import numpy as np
1011

@@ -16,6 +17,7 @@
1617
"get_total_and_free_memory_in_Mb",
1718
"occupy_mem",
1819
"gpu_mem_usage",
20+
"mem_usage"
1921
]
2022

2123

@@ -51,6 +53,15 @@ def gpu_mem_usage():
5153
return mem_usage_bytes / (1024 * 1024)
5254

5355

56+
def mem_usage():
57+
"""
58+
Compute the memory usage for the current machine (GB).
59+
"""
60+
gb = 1 << 30
61+
mem = psutil.virtual_memory()
62+
return mem.used / gb
63+
64+
5465
class AverageMeter:
5566
"""Track a series of values and provide access to smoothed values over a
5667
window or the global series average.

0 commit comments

Comments
 (0)