Skip to content

Commit 7a62f48

Browse files
repair_some_Bug_for_paddlex
1 parent 6dc0211 commit 7a62f48

File tree

10 files changed

+49
-87
lines changed

10 files changed

+49
-87
lines changed

docs/algorithm/formula_recognition/algorithm_rec_latex_ocr.en.md

+6-6
Original file line numberDiff line numberDiff line change
@@ -17,7 +17,7 @@ Please refer to ["Environment Preparation"](../../ppocr/environment.en.md) to co
1717

1818
Furthermore, additional dependencies need to be installed:
1919
```shell
20-
pip install "tokenizers==0.19.1" "imagesize"
20+
pip install -r docs/algorithm/formula_recognition/requirements.txt
2121
```
2222

2323
## 3. Model Training / Evaluation / Prediction
@@ -61,16 +61,16 @@ Evaluation:
6161
```
6262
# GPU evaluation
6363
# Validation set evaluation
64-
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Metric.cal_blue_score=True
64+
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams
6565
# Test set evaluation
66-
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Metric.cal_blue_score=True Eval.dataset.data=./train_data/LaTeXOCR/latexocr_test.pkl
66+
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Eval.dataset.data=./train_data/LaTeXOCR/latexocr_test.pkl
6767
```
6868

6969
Prediction:
7070

7171
```
7272
# The configuration file used for prediction must match the training
73-
python3 tools/infer_rec.py -c configs/rec/rec_latex_ocr.yml -o Architecture.Backbone.is_predict=True Architecture.Backbone.is_export=True Architecture.Head.is_export=True Global.infer_img='./doc/datasets/pme_demo/0000013.png' Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams
73+
python3 tools/infer_rec.py -c configs/rec/rec_latex_ocr.yml -o Global.infer_img='./docs/datasets/images/pme_demo/0000013.png' Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams
7474
```
7575

7676
## 4. Inference and Deployment
@@ -79,15 +79,15 @@ python3 tools/infer_rec.py -c configs/rec/rec_latex_ocr.yml -o Architecture.Ba
7979
First, the model saved during the LaTeX-OCR printed mathematical expression recognition training process is converted into an inference model. you can use the following command to convert:
8080

8181
```
82-
python3 tools/export_model.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Global.save_inference_dir=./inference/rec_latex_ocr_infer/ Architecture.Backbone.is_predict=True Architecture.Backbone.is_export=True Architecture.Head.is_export=True
82+
python3 tools/export_model.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Global.save_inference_dir=./inference/rec_latex_ocr_infer/
8383
8484
# The default output max length of the model is 512.
8585
```
8686

8787
For LaTeX-OCR printed mathematical expression recognition model inference, the following commands can be executed:
8888

8989
```
90-
python3 tools/infer/predict_rec.py --image_dir='./doc/datasets/pme_demo/0000295.png' --rec_algorithm="LaTeXOCR" --rec_batch_num=1 --rec_model_dir="./inference/rec_latex_ocr_infer/" --rec_char_dict_path="./ppocr/utils/dict/latex_ocr_tokenizer.json"
90+
python3 tools/infer/predict_rec.py --image_dir='./docs/datasets/images/pme_demo/0000295.png' --rec_algorithm="LaTeXOCR" --rec_batch_num=1 --rec_model_dir="./inference/rec_latex_ocr_infer/" --rec_char_dict_path="./ppocr/utils/dict/latex_ocr_tokenizer.json"
9191
```
9292

9393
### 4.2 C++ Inference

docs/algorithm/formula_recognition/algorithm_rec_latex_ocr.md

+6-6
Original file line numberDiff line numberDiff line change
@@ -17,7 +17,7 @@
1717

1818
此外,需要安装额外的依赖:
1919
```shell
20-
pip install "tokenizers==0.19.1" "imagesize"
20+
pip install -r docs/algorithm/formula_recognition/requirements.txt
2121
```
2222

2323
## 3. 模型训练、评估、预测
@@ -69,17 +69,17 @@ python3 tools/train.py -c configs/rec/rec_latex_ocr.yml -o Global.eval_batch_ste
6969
```shell
7070
# 注意将pretrained_model的路径设置为本地路径。若使用自行训练保存的模型,请注意修改路径和文件名为{path/to/weights}/{model_name}。
7171
# 验证集评估
72-
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Metric.cal_blue_score=True
72+
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams
7373
# 测试集评估
74-
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Metric.cal_blue_score=True Eval.dataset.data=./train_data/LaTeXOCR/latexocr_test.pkl
74+
python3 tools/eval.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Eval.dataset.data=./train_data/LaTeXOCR/latexocr_test.pkl
7575
```
7676

7777
### 3.4 预测
7878

7979
使用如下命令进行单张图片预测:
8080
```shell
8181
# 注意将pretrained_model的路径设置为本地路径。
82-
python3 tools/infer_rec.py -c configs/rec/rec_latex_ocr.yml -o Architecture.Backbone.is_predict=True Architecture.Backbone.is_export=True Architecture.Head.is_export=True Global.infer_img='./doc/datasets/pme_demo/0000013.png' Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams
82+
python3 tools/infer_rec.py -c configs/rec/rec_latex_ocr.yml -o Global.infer_img='./docs/datasets/images/pme_demo/0000013.png' Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams
8383
# 预测文件夹下所有图像时,可修改infer_img为文件夹,如 Global.infer_img='./doc/datasets/pme_demo/'。
8484
```
8585

@@ -90,7 +90,7 @@ python3 tools/infer_rec.py -c configs/rec/rec_latex_ocr.yml -o Architecture.Ba
9090

9191
```shell
9292
# 注意将pretrained_model的路径设置为本地路径。
93-
python3 tools/export_model.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Global.save_inference_dir=./inference/rec_latex_ocr_infer/ Architecture.Backbone.is_predict=True Architecture.Backbone.is_export=True Architecture.Head.is_export=True
93+
python3 tools/export_model.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrained_model=./rec_latex_ocr_train/best_accuracy.pdparams Global.save_inference_dir=./inference/rec_latex_ocr_infer/
9494

9595
# 目前的静态图模型支持的最大输出长度为512
9696
```
@@ -109,7 +109,7 @@ python3 tools/export_model.py -c configs/rec/rec_latex_ocr.yml -o Global.pretrai
109109
执行如下命令进行模型推理:
110110

111111
```shell
112-
python3 tools/infer/predict_rec.py --image_dir='./doc/datasets/pme_demo/0000295.png' --rec_algorithm="LaTeXOCR" --rec_batch_num=1 --rec_model_dir="./inference/rec_latex_ocr_infer/" --rec_char_dict_path="./ppocr/utils/dict/latex_ocr_tokenizer.json"
112+
python3 tools/infer/predict_rec.py --image_dir='./docs/datasets/images/pme_demo/0000295.png' --rec_algorithm="LaTeXOCR" --rec_batch_num=1 --rec_model_dir="./inference/rec_latex_ocr_infer/" --rec_char_dict_path="./ppocr/utils/dict/latex_ocr_tokenizer.json"
113113

114114
# 预测文件夹下所有图像时,可修改image_dir为文件夹,如 --image_dir='./doc/datasets/pme_demo/'。
115115
```
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,2 @@
1+
tokenizers==0.19.1
2+
imagesize

ppocr/metrics/rec_metric.py

+3-3
Original file line numberDiff line numberDiff line change
@@ -261,9 +261,9 @@ def get_metric(self):
261261
self.reset()
262262
if self.cal_blue_score:
263263
return {
264-
"blue_score ": cur_blue_score,
265-
"edit distance ": cur_edit_distance,
266-
"exp_rate ": cur_exp_rate,
264+
"blue_score": cur_blue_score,
265+
"edit distance": cur_edit_distance,
266+
"exp_rate": cur_exp_rate,
267267
"exp_rate<=1 ": cur_exp_1,
268268
"exp_rate<=2 ": cur_exp_2,
269269
"exp_rate<=3 ": cur_exp_3,

ppocr/modeling/backbones/rec_resnetv2.py

+1-62
Original file line numberDiff line numberDiff line change
@@ -546,65 +546,6 @@ def forward(self, x):
546546
check_variable_and_dtype,
547547
)
548548

549-
550-
def group_norm(
551-
input,
552-
groups,
553-
epsilon=1e-05,
554-
weight=None,
555-
bias=None,
556-
act=None,
557-
data_layout="NCHW",
558-
name=None,
559-
):
560-
helper = LayerHelper("group_norm", **locals())
561-
dtype = helper.input_dtype()
562-
check_variable_and_dtype(
563-
input,
564-
"input",
565-
["float16", "uint16", "float32", "float64"],
566-
"group_norm",
567-
)
568-
# create intput and parameters
569-
inputs = {"X": input}
570-
input_shape = input.shape
571-
if len(input_shape) < 2:
572-
raise ValueError(
573-
f"The dimensions of Op(static.nn.group_norm)'s input should be more than 1. But received {len(input_shape)}"
574-
)
575-
if data_layout != "NCHW" and data_layout != "NHWC":
576-
raise ValueError(
577-
"Param(data_layout) of Op(static.nn.group_norm) got wrong value: received "
578-
+ data_layout
579-
+ " but only NCHW or NHWC supported."
580-
)
581-
channel_num = input_shape[1] if data_layout == "NCHW" else input_shape[-1]
582-
param_shape = [channel_num]
583-
inputs["Scale"] = weight
584-
inputs["Bias"] = bias
585-
# create output
586-
mean_out = helper.create_variable(dtype=dtype, stop_gradient=True)
587-
variance_out = helper.create_variable(dtype=dtype, stop_gradient=True)
588-
group_norm_out = helper.create_variable(dtype=dtype)
589-
590-
helper.append_op(
591-
type="group_norm",
592-
inputs=inputs,
593-
outputs={
594-
"Y": group_norm_out,
595-
"Mean": mean_out,
596-
"Variance": variance_out,
597-
},
598-
attrs={
599-
"epsilon": epsilon,
600-
"groups": groups,
601-
"data_layout": data_layout,
602-
},
603-
)
604-
605-
return helper.append_activation(group_norm_out)
606-
607-
608549
class GroupNormAct(nn.GroupNorm):
609550
# NOTE num_channel and num_groups order flipped for easier layer swaps / binding of fixed args
610551
def __init__(
@@ -630,9 +571,7 @@ def __init__(
630571
self.act = nn.Identity()
631572

632573
def forward(self, x):
633-
x = group_norm(
634-
x, self._num_groups, self._epsilon, weight=self.weight, bias=self.bias
635-
)
574+
x = F.group_norm(x, num_groups=self._num_groups, weight=self.weight, bias=self.bias)
636575
x = self.act(x)
637576
return x
638577

ppocr/modeling/heads/rec_latexocr_head.py

+1-4
Original file line numberDiff line numberDiff line change
@@ -977,10 +977,7 @@ def generate_export(
977977
paddle.cumsum((out == eos_token).cast(paddle.int64), 1)[:, -1] >= 1
978978
).all()
979979
):
980-
out = out[:, t:]
981-
if num_dims == 1:
982-
out = out.squeeze(0)
983-
return out
980+
break
984981
i_idx += 1
985982
out = out[:, t:]
986983
if num_dims == 1:

ppocr/utils/formula_utils/math_txt2pkl.py

+4-1
Original file line numberDiff line numberDiff line change
@@ -15,6 +15,7 @@
1515
import pickle
1616
from tqdm import tqdm
1717
import os
18+
import math
1819
from paddle.utils import try_import
1920
from collections import defaultdict
2021
import glob
@@ -41,7 +42,9 @@ def txt2pickle(images, equations, save_dir):
4142
min_dimensions[0] <= width <= max_dimensions[0]
4243
and min_dimensions[1] <= height <= max_dimensions[1]
4344
):
44-
data[(width, height)].append((eqs[indices[i]], im))
45+
divide_h = math.ceil(height / 16) * 16
46+
divide_w = math.ceil(width / 16) * 16
47+
data[(divide_w, divide_h)].append((eqs[indices[i]], im))
4548
data = dict(data)
4649
with open(save_p, "wb") as file:
4750
pickle.dump(data, file)

tools/eval.py

+1
Original file line numberDiff line numberDiff line change
@@ -107,6 +107,7 @@ def main():
107107
model_type = "can"
108108
elif config["Architecture"]["algorithm"] == "LaTeXOCR":
109109
model_type = "latexocr"
110+
config["Metric"]["cal_blue_score"] = True
110111
else:
111112
model_type = config["Architecture"]["model_type"]
112113
else:

tools/export_model.py

+19-5
Original file line numberDiff line numberDiff line change
@@ -22,6 +22,7 @@
2222
import argparse
2323

2424
import yaml
25+
import json
2526
import paddle
2627
from paddle.jit import to_static
2728
from collections import OrderedDict
@@ -219,11 +220,18 @@ def dump_infer_config(config, path, logger):
219220
for k, v in config["PostProcess"].items():
220221
postprocess[k] = v
221222

222-
if config["Global"].get("character_dict_path") is not None:
223-
with open(config["Global"]["character_dict_path"], encoding="utf-8") as f:
224-
lines = f.readlines()
225-
character_dict = [line.strip("\n") for line in lines]
226-
postprocess["character_dict"] = character_dict
223+
if config["Architecture"].get("algorithm") in ["LaTeXOCR"]:
224+
tokenizer_file = config["Global"].get("rec_char_dict_path")
225+
if tokenizer_file is not None:
226+
with open(tokenizer_file, encoding="utf-8") as tokenizer_config_handle:
227+
character_dict = json.load(tokenizer_config_handle)
228+
postprocess["character_dict"] = character_dict
229+
else:
230+
if config["Global"].get("character_dict_path") is not None:
231+
with open(config["Global"]["character_dict_path"], encoding="utf-8") as f:
232+
lines = f.readlines()
233+
character_dict = [line.strip("\n") for line in lines]
234+
postprocess["character_dict"] = character_dict
227235

228236
infer_cfg["PostProcess"] = postprocess
229237

@@ -288,6 +296,12 @@ def main():
288296
# for sr algorithm
289297
if config["Architecture"]["model_type"] == "sr":
290298
config["Architecture"]["Transform"]["infer_mode"] = True
299+
300+
# for latexocr algorithm
301+
if config["Architecture"].get("algorithm") in ["LaTeXOCR"]:
302+
config["Architecture"]["Backbone"]["is_predict"] = True
303+
config["Architecture"]["Backbone"]["is_export"] = True
304+
config["Architecture"]["Head"]["is_export"] = True
291305
model = build_model(config["Architecture"])
292306
load_model(config, model, model_type=config["Architecture"]["model_type"])
293307
model.eval()

tools/infer_rec.py

+6
Original file line numberDiff line numberDiff line change
@@ -82,6 +82,12 @@ def main():
8282
config["Architecture"]["Head"]["out_channels_list"] = out_channels_list
8383
else: # base rec model
8484
config["Architecture"]["Head"]["out_channels"] = char_num
85+
86+
if config["Architecture"].get("algorithm") in ["LaTeXOCR"]:
87+
config["Architecture"]["Backbone"]["is_predict"] = True
88+
config["Architecture"]["Backbone"]["is_export"] = True
89+
config["Architecture"]["Head"]["is_export"] = True
90+
8591
model = build_model(config["Architecture"])
8692

8793
load_model(config, model)

0 commit comments

Comments
 (0)